Skip to content

Commit 35f3e9d

Browse files
committed
Merge commit for internal changes
2 parents 7d57bf4 + 04a318d commit 35f3e9d

22 files changed

+2072
-439
lines changed

docs/tutorials/shape_constraints.ipynb

Lines changed: 8 additions & 8 deletions
Original file line numberDiff line numberDiff line change
@@ -154,9 +154,9 @@
154154
},
155155
"outputs": [],
156156
"source": [
157-
"NUM_EPOCHS = 500\n",
157+
"NUM_EPOCHS = 1000\n",
158158
"BATCH_SIZE = 64\n",
159-
"LEARNING_RATE=0.001"
159+
"LEARNING_RATE=0.01"
160160
]
161161
},
162162
{
@@ -376,9 +376,9 @@
376376
"\n",
377377
"# Generate datasets.\n",
378378
"np.random.seed(42)\n",
379-
"data_train = sample_dataset(2000, testing_set=False)\n",
380-
"data_val = sample_dataset(1000, testing_set=False)\n",
381-
"data_test = sample_dataset(1000, testing_set=True)\n",
379+
"data_train = sample_dataset(500, testing_set=False)\n",
380+
"data_val = sample_dataset(500, testing_set=False)\n",
381+
"data_test = sample_dataset(500, testing_set=True)\n",
382382
"\n",
383383
"# Plotting dataset densities.\n",
384384
"figsize(12, 5)\n",
@@ -536,9 +536,9 @@
536536
" feature_columns=feature_columns,\n",
537537
" # Hyper-params optimized on validation set.\n",
538538
" n_batches_per_layer=1,\n",
539-
" max_depth=3,\n",
540-
" n_trees=20,\n",
541-
" min_node_weight=0.1,\n",
539+
" max_depth=2,\n",
540+
" n_trees=50,\n",
541+
" learning_rate=0.05,\n",
542542
" config=tf.estimator.RunConfig(tf_random_seed=42),\n",
543543
")\n",
544544
"gbt_estimator.train(input_fn=train_input_fn)\n",

examples/custom_estimators_uci_heart.py

Lines changed: 13 additions & 6 deletions
Original file line numberDiff line numberDiff line change
@@ -18,7 +18,7 @@
1818
This example trains a TFL custom estimators on the UCI heart dataset.
1919
2020
Example usage:
21-
custom_estimators_uci_heart --num_epochs=40
21+
custom_estimators_uci_heart --num_epochs=5000
2222
"""
2323

2424
from __future__ import absolute_import
@@ -38,7 +38,7 @@
3838
FLAGS = flags.FLAGS
3939
flags.DEFINE_float('learning_rate', 0.01, 'Learning rate.')
4040
flags.DEFINE_integer('batch_size', 100, 'Batch size.')
41-
flags.DEFINE_integer('num_epochs', 200, 'Number of training epoch.')
41+
flags.DEFINE_integer('num_epochs', 2000, 'Number of training epoch.')
4242

4343

4444
def main(_):
@@ -121,7 +121,7 @@ def model_fn(features, labels, mode, config):
121121
tfl.layers.PWLCalibration(
122122
input_keypoints=[0.0, 1.0, 2.0, 3.0],
123123
output_min=0.0,
124-
output_max=lattice_sizes[0] - 1.0,
124+
output_max=lattice_sizes[2] - 1.0,
125125
# You can specify TFL regularizers as tuple
126126
# ('regularizer name', l1, l2).
127127
kernel_regularizer=('hessian', 0.0, 1e-4),
@@ -130,16 +130,23 @@ def model_fn(features, labels, mode, config):
130130
tfl.layers.CategoricalCalibration(
131131
num_buckets=3,
132132
output_min=0.0,
133-
output_max=lattice_sizes[1] - 1.0,
133+
output_max=lattice_sizes[3] - 1.0,
134134
# Categorical monotonicity can be partial order.
135135
# (i, j) indicates that we must have output(i) <= output(i).
136136
# Make sure to set the lattice monotonicity to 1 for this dimension.
137137
monotonicities=[(0, 1), (0, 2)],
138138
)(inputs['thal']),
139139
])
140140
output = tfl.layers.Lattice(
141-
lattice_sizes=lattice_sizes, monotonicities=lattice_monotonicities)(
142-
lattice_input)
141+
lattice_sizes=lattice_sizes,
142+
monotonicities=lattice_monotonicities,
143+
# Add a kernel_initializer so that the Lattice is not initialized as a
144+
# flat plane. The output_min and output_max could be arbitrary, as long
145+
# as output_min < output_max.
146+
kernel_initializer=tfl.lattice_layer.RandomMonotonicInitializer(
147+
lattice_sizes=lattice_sizes, output_min=-10, output_max=10),
148+
)(
149+
lattice_input)
143150

144151
training = (mode == tf.estimator.ModeKeys.TRAIN)
145152
model = tf.keras.Model(inputs=inputs, outputs=output)

setup.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -27,7 +27,7 @@
2727
# This version number should always be that of the *next* (unreleased) version.
2828
# Immediately after uploading a package to PyPI, you should increment the
2929
# version number and push to gitHub.
30-
__version__ = "2.0.7"
30+
__version__ = "2.0.8"
3131

3232
if "--release" in sys.argv:
3333
sys.argv.remove("--release")

tensorflow_lattice/python/BUILD

Lines changed: 7 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -117,6 +117,7 @@ py_library(
117117
deps = [
118118
":categorical_calibration_layer",
119119
":configs",
120+
":kronecker_factored_lattice_layer",
120121
":lattice_layer",
121122
":linear_layer",
122123
":model_info",
@@ -197,6 +198,7 @@ py_test(
197198
srcs_version = "PY2AND3",
198199
deps = [
199200
":kronecker_factored_lattice_layer",
201+
":kronecker_factored_lattice_lib",
200202
":test_utils",
201203
# absl/logging dep,
202204
# absl/testing:parameterized dep,
@@ -329,6 +331,7 @@ py_library(
329331
":aggregation_layer",
330332
":categorical_calibration_layer",
331333
":configs",
334+
":kronecker_factored_lattice_layer",
332335
":lattice_layer",
333336
":parallel_combination_layer",
334337
":premade_lib",
@@ -346,6 +349,8 @@ py_library(
346349
":aggregation_layer",
347350
":categorical_calibration_layer",
348351
":configs",
352+
":kronecker_factored_lattice_layer",
353+
":kronecker_factored_lattice_lib",
349354
":lattice_layer",
350355
":lattice_lib",
351356
":linear_layer",
@@ -372,6 +377,7 @@ py_test(
372377
":premade",
373378
":premade_lib",
374379
# absl/logging dep,
380+
# absl/testing:parameterized dep,
375381
# numpy dep,
376382
# tensorflow dep,
377383
],
@@ -439,6 +445,7 @@ py_library(
439445
srcs = ["rtl_layer.py"],
440446
srcs_version = "PY2AND3",
441447
deps = [
448+
":kronecker_factored_lattice_layer",
442449
":lattice_layer",
443450
":rtl_lib",
444451
# tensorflow:tensorflow_no_contrib dep,

tensorflow_lattice/python/configs.py

Lines changed: 65 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -267,6 +267,8 @@ def __init__(self,
267267
num_lattices=None,
268268
lattice_rank=None,
269269
interpolation='hypercube',
270+
parameterization='all_vertices',
271+
num_terms=2,
270272
separate_calibrators=True,
271273
use_linear_combination=False,
272274
use_bias=False,
@@ -305,6 +307,34 @@ def __init__(self,
305307
'simplex' uses d+1 parameters and thus scales better. For details see
306308
`tfl.lattice_lib.evaluate_with_simplex_interpolation` and
307309
`tfl.lattice_lib.evaluate_with_hypercube_interpolation`.
310+
parameterization: The parameterization of the lattice function class to
311+
use. A lattice function is uniquely determined by specifying its value
312+
on every lattice vertex. A parameterization scheme is a mapping from a
313+
vector of parameters to a multidimensional array of lattice vertex
314+
values. It can be one of:
315+
- String `'all_vertices'`: This is the "traditional" parameterization
316+
that keeps one scalar parameter per lattice vertex where the mapping
317+
is essentially the identity map. With this scheme, the number of
318+
parameters scales exponentially with the number of inputs to the
319+
lattice. The underlying lattices used will be `tfl.layers.Lattice`
320+
layers.
321+
- String `'kronecker_factored'`: With this parameterization, for each
322+
lattice input i we keep a collection of `num_terms` vectors each
323+
having `feature_configs[0].lattice_size` entries (note that all
324+
features must have the same lattice size). To obtain the tensor of
325+
lattice vertex values, for `t=1,2,...,num_terms` we compute the
326+
outer product of the `t'th` vector in each collection, multiply by a
327+
per-term scale, and sum the resulting tensors. Finally, we add a
328+
single shared bias parameter to each entry in the sum. With this
329+
scheme, the number of parameters grows linearly with `lattice_rank`
330+
(assuming lattice sizes and `num_terms` are held constant).
331+
Currently, only monotonicity shape constraint and bound constraint
332+
are supported for this scheme. Regularization is not currently
333+
supported. The underlying lattices used will be
334+
`tfl.layers.KroneckerFactoredLattice` layers.
335+
num_terms: The number of terms in a lattice using `'kronecker_factored'`
336+
parameterization. Ignored if parameterization is set to
337+
`'all_vertices'`.
308338
separate_calibrators: If features should be separately calibrated for each
309339
lattice in the ensemble.
310340
use_linear_combination: If set to true, a linear combination layer will be
@@ -375,12 +405,15 @@ class CalibratedLatticeConfig(_Config, _HasFeatureConfigs,
375405
def __init__(self,
376406
feature_configs=None,
377407
interpolation='hypercube',
408+
parameterization='all_vertices',
409+
num_terms=2,
378410
regularizer_configs=None,
379411
output_min=None,
380412
output_max=None,
381413
output_calibration=False,
382414
output_calibration_num_keypoints=10,
383-
output_initialization='quantiles'):
415+
output_initialization='quantiles',
416+
random_seed=0):
384417
"""Initializes a `CalibratedLatticeConfig` instance.
385418
386419
Args:
@@ -392,6 +425,34 @@ def __init__(self,
392425
'simplex' uses d+1 parameters and thus scales better. For details see
393426
`tfl.lattice_lib.evaluate_with_simplex_interpolation` and
394427
`tfl.lattice_lib.evaluate_with_hypercube_interpolation`.
428+
parameterization: The parameterization of the lattice function class to
429+
use. A lattice function is uniquely determined by specifying its value
430+
on every lattice vertex. A parameterization scheme is a mapping from a
431+
vector of parameters to a multidimensional array of lattice vertex
432+
values. It can be one of:
433+
- String `'all_vertices'`: This is the "traditional" parameterization
434+
that keeps one scalar parameter per lattice vertex where the mapping
435+
is essentially the identity map. With this scheme, the number of
436+
parameters scales exponentially with the number of inputs to the
437+
lattice. The underlying lattice used will be a `tfl.layers.Lattice`
438+
layer.
439+
- String `'kronecker_factored'`: With this parameterization, for each
440+
lattice input i we keep a collection of `num_terms` vectors each
441+
having `feature_configs[0].lattice_size` entries (note that all
442+
features must have the same lattice size). To obtain the tensor of
443+
lattice vertex values, for `t=1,2,...,num_terms` we compute the
444+
outer product of the `t'th` vector in each collection, multiply by a
445+
per-term scale, and sum the resulting tensors. Finally, we add a
446+
single shared bias parameter to each entry in the sum. With this
447+
scheme, the number of parameters grows linearly with
448+
`len(feature_configs)` (assuming lattice sizes and `num_terms` are
449+
held constant). Currently, only monotonicity shape constraint and
450+
bound constraint are supported for this scheme. Regularization is
451+
not currently supported. The underlying lattice used will be a
452+
`tfl.layers.KroneckerFactoredLattice` layer.
453+
num_terms: The number of terms in a lattice using `'kronecker_factored'`
454+
parameterization. Ignored if parameterization is set to
455+
`'all_vertices'`.
395456
regularizer_configs: A list of `tfl.configs.RegularizerConfig` instances
396457
that apply global regularization.
397458
output_min: Lower bound constraint on the output of the model.
@@ -410,6 +471,9 @@ def __init__(self,
410471
- String `'uniform'`: Output is initliazed uniformly in label range.
411472
- A list of numbers: To be used for initialization of the output
412473
lattice or output calibrator.
474+
random_seed: Random seed to use for initialization of a lattice with
475+
`'kronecker_factored'` parameterization. Ignored if parameterization is
476+
set to `'all_vertices'`.
413477
"""
414478
super(CalibratedLatticeConfig, self).__init__(locals())
415479

0 commit comments

Comments
 (0)