Skip to content
This repository was archived by the owner on Apr 10, 2024. It is now read-only.

Commit e830c8f

Browse files
Increment version to 0.1.0
1 parent eee40f3 commit e830c8f

File tree

8 files changed

+119
-12
lines changed

8 files changed

+119
-12
lines changed

README.md

Lines changed: 17 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -36,7 +36,7 @@ run in your browser.
3636
<img src="https://storage.googleapis.com/lucid-static/common/stickers/colab-tutorial.png" width="500" alt=""></img>
3737
</a>
3838

39-
## Building Blocks
39+
## Building Blocks
4040
*Notebooks corresponding to the [Building Blocks of Interpretability](https://distill.pub/2018/building-blocks/) article*
4141

4242

@@ -80,6 +80,22 @@ This project is research code. It is not an official Google product.
8080

8181
## Development
8282

83+
### Style guide deviations
84+
85+
We use naming conventions to help differentiate tensors, operations, and values:
86+
87+
* Suffix variable names representing **tensors** with `_t`
88+
* Suffix variable names representing **operations** with `_op`
89+
* Don't suffix variable names representing concrete values
90+
91+
Usage example:
92+
93+
```
94+
global_step_t = tf.train.get_or_create_global_step()
95+
global_step_init_op = tf.variables_initializer([global_step_t])
96+
global_step = global_step_t.eval()
97+
```
98+
8399
### Running Tests
84100

85101
Use `tox` to run the test suite on all supported environments.

lucid/misc/io/serialize_array.py

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -48,6 +48,7 @@ def _normalize_array(array, domain=(0, 1)):
4848
array = np.squeeze(array)
4949
assert len(array.shape) <= 3
5050
assert np.issubdtype(array.dtype, np.number)
51+
assert not np.isnan(array).any()
5152

5253
low, high = np.min(array), np.max(array)
5354
if domain is None:

lucid/misc/redirected_relu_grad.py

Lines changed: 20 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -23,13 +23,20 @@
2323
These functions provide a more convenient solution: temporarily override the
2424
gradient of ReLUs to allow gradient to flow back through the ReLU -- even if it
2525
didn't activate and had a derivative of zero -- allowing the visualization
26-
process to get started.
26+
process to get started. These functions override the gradient for at most 16
27+
steps. Thus, you need to initialize `global_step` before using these functions.
2728
2829
Usage:
2930
```python
3031
from lucid.misc.gradient_override import gradient_override_map
3132
from lucid.misc.redirected_relu_grad import redirected_relu_grad
3233
34+
...
35+
global_step_t = tf.train.get_or_create_global_step()
36+
init_global_step_op = tf.variables_initializer([global_step_t])
37+
init_global_step_op.run()
38+
...
39+
3340
with gradient_override_map({'Relu': redirected_relu_grad}):
3441
model.import_graph(...)
3542
```
@@ -99,7 +106,12 @@ def redirected_relu_grad(op, grad):
99106
batch = tf.shape(relu_grad)[0]
100107
reshaped_relu_grad = tf.reshape(relu_grad, [batch, -1])
101108
relu_grad_mag = tf.norm(reshaped_relu_grad, axis=1)
102-
return tf.where(relu_grad_mag > 0., relu_grad, redirected_grad)
109+
result_grad = tf.where(relu_grad_mag > 0., relu_grad, redirected_grad)
110+
111+
global_step_t =tf.train.get_or_create_global_step()
112+
return_relu_grad = tf.greater(global_step_t, tf.constant(16, tf.int64))
113+
114+
return tf.where(return_relu_grad, relu_grad, result_grad)
103115

104116

105117
def redirected_relu6_grad(op, grad):
@@ -125,4 +137,9 @@ def redirected_relu6_grad(op, grad):
125137
batch = tf.shape(relu_grad)[0]
126138
reshaped_relu_grad = tf.reshape(relu_grad, [batch, -1])
127139
relu_grad_mag = tf.norm(reshaped_relu_grad, axis=1)
128-
return tf.where(relu_grad_mag > 0., relu_grad, redirected_grad)
140+
result_grad = tf.where(relu_grad_mag > 0., relu_grad, redirected_grad)
141+
142+
global_step_t = tf.train.get_or_create_global_step()
143+
return_relu_grad = tf.greater(global_step_t, tf.constant(16, tf.int64))
144+
145+
return tf.where(return_relu_grad, relu_grad, result_grad)

lucid/optvis/objectives.py

Lines changed: 17 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -131,7 +131,23 @@ def wrap_objective(f, *args, **kwds):
131131

132132
@wrap_objective
133133
def neuron(layer_name, channel_n, x=None, y=None, batch=None):
134-
"""Visualize a single neuron of a single channel."""
134+
"""Visualize a single neuron of a single channel.
135+
136+
Defaults to the center neuron. When width and height are even numbers, we
137+
choose the neuron in the bottom right of the center 2x2 neurons.
138+
139+
Odd width & height: Even width & height:
140+
141+
+---+---+---+ +---+---+---+---+
142+
| | | | | | | | |
143+
+---+---+---+ +---+---+---+---+
144+
| | X | | | | | | |
145+
+---+---+---+ +---+---+---+---+
146+
| | | | | | | X | |
147+
+---+---+---+ +---+---+---+---+
148+
| | | | |
149+
+---+---+---+---+
150+
"""
135151
def inner(T):
136152
layer = T(layer_name)
137153
shape = tf.shape(layer)

lucid/optvis/render.py

Lines changed: 26 additions & 6 deletions
Original file line numberDiff line numberDiff line change
@@ -31,6 +31,8 @@
3131

3232
from lucid.optvis import objectives, param, transform
3333
from lucid.misc.io import show
34+
from lucid.misc.redirected_relu_grad import redirected_relu_grad, redirected_relu6_grad
35+
from lucid.misc.gradient_override import gradient_override_map
3436

3537
# pylint: disable=invalid-name
3638

@@ -40,8 +42,8 @@
4042

4143

4244
def render_vis(model, objective_f, param_f=None, optimizer=None,
43-
transforms=None, thresholds=(512,),
44-
print_objectives=None, verbose=True,):
45+
transforms=None, thresholds=(512,), print_objectives=None,
46+
verbose=True, relu_gradient_override=True, use_fixed_seed=False):
4547
"""Flexible optimization-base feature vis.
4648
4749
There's a lot of ways one might wish to customize otpimization-based
@@ -72,6 +74,11 @@ def render_vis(model, objective_f, param_f=None, optimizer=None,
7274
whose values get logged during the optimization.
7375
verbose: Should we display the visualization when we hit a threshold?
7476
This should only be used in IPython.
77+
relu_gradient_override: Whether to use the gradient override scheme
78+
described in lucid/misc/redirected_relu_grad.py. On by default!
79+
use_fixed_seed: Seed the RNG with a fixed value so results are reproducible.
80+
Off by default. As of tf 1.8 this does not work as intended, see:
81+
https://github.com/tensorflow/tensorflow/issues/9171
7582
Returns:
7683
2D array of optimization results containing of evaluations of supplied
7784
param_f snapshotted at specified thresholds. Usually that will mean one or
@@ -80,7 +87,11 @@ def render_vis(model, objective_f, param_f=None, optimizer=None,
8087

8188
with tf.Graph().as_default() as graph, tf.Session() as sess:
8289

83-
T = make_vis_T(model, objective_f, param_f, optimizer, transforms)
90+
if use_fixed_seed: # does not mean results are reproducible, see Args doc
91+
tf.set_random_seed(0)
92+
93+
T = make_vis_T(model, objective_f, param_f, optimizer, transforms,
94+
relu_gradient_override)
8495
print_objective_func = make_print_objective_func(print_objectives, T)
8596
loss, vis_op, t_image = T("loss"), T("vis_op"), T("input")
8697
tf.global_variables_initializer().run()
@@ -105,7 +116,7 @@ def render_vis(model, objective_f, param_f=None, optimizer=None,
105116

106117

107118
def make_vis_T(model, objective_f, param_f=None, optimizer=None,
108-
transforms=None):
119+
transforms=None, relu_gradient_override=False):
109120
"""Even more flexible optimization-base feature vis.
110121
111122
This function is the inner core of render_vis(), and can be used
@@ -155,10 +166,19 @@ def make_vis_T(model, objective_f, param_f=None, optimizer=None,
155166
transform_f = make_transform_f(transforms)
156167
optimizer = make_optimizer(optimizer, [])
157168

158-
T = import_model(model, transform_f(t_image), t_image)
169+
global_step = tf.train.get_or_create_global_step()
170+
init_global_step = tf.variables_initializer([global_step])
171+
init_global_step.run()
172+
173+
if relu_gradient_override:
174+
with gradient_override_map({'Relu': redirected_relu_grad,
175+
'Relu6': redirected_relu6_grad}):
176+
T = import_model(model, transform_f(t_image), t_image)
177+
else:
178+
T = import_model(model, transform_f(t_image), t_image)
159179
loss = objective_f(T)
160180

161-
global_step = tf.Variable(0, trainable=False, name="global_step")
181+
162182
vis_op = optimizer.minimize(-loss, global_step=global_step)
163183

164184
local_vars = locals()

setup.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -18,7 +18,7 @@
1818

1919
from setuptools import setup, find_packages
2020

21-
version = '0.0.8'
21+
version = '0.1.0'
2222

2323
test_deps = [
2424
'future',

tests/misc/test_gradient_override.py

Lines changed: 32 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -29,6 +29,10 @@ def gradient_override(op, grad):
2929
return tf.constant(42)
3030

3131
with tf.Session().as_default() as sess:
32+
global_step = tf.train.get_or_create_global_step()
33+
init_global_step = tf.variables_initializer([global_step])
34+
init_global_step.run()
35+
3236
a = tf.constant(1.)
3337
standard_relu = tf.nn.relu(a)
3438
grad_wrt_a = tf.gradients(standard_relu, a, [1.])[0]
@@ -56,10 +60,38 @@ def test_gradient_override_relu6_directionality(nonl_name, nonl,
5660
nonl_grad_override, examples):
5761
for incoming_grad, input, grad in examples:
5862
with tf.Session().as_default() as sess:
63+
global_step = tf.train.get_or_create_global_step()
64+
init_global_step = tf.variables_initializer([global_step])
65+
init_global_step.run()
66+
5967
batched_shape = [1,1]
6068
incoming_grad_t = tf.constant(incoming_grad, shape=batched_shape)
6169
input_t = tf.constant(input, shape=batched_shape)
6270
with gradient_override_map({nonl_name: nonl_grad_override}):
6371
nonl_t = nonl(input_t)
6472
grad_wrt_input = tf.gradients(nonl_t, input_t, [incoming_grad_t])[0]
6573
assert (grad_wrt_input.eval() == grad).all()
74+
75+
@pytest.mark.parametrize("nonl_name,nonl,nonl_grad_override, examples", nonls)
76+
def test_gradient_override_shutoff(nonl_name, nonl,
77+
nonl_grad_override, examples):
78+
for incoming_grad, input, grad in examples:
79+
with tf.Session().as_default() as sess:
80+
global_step_t = tf.train.get_or_create_global_step()
81+
global_step_init_op = tf.variables_initializer([global_step_t])
82+
global_step_init_op.run()
83+
global_step_assign_t = tf.assign(global_step_t, 17)
84+
sess.run(global_step_assign_t)
85+
86+
# similar setup to test_gradient_override_relu6_directionality,
87+
# but we test that the gradient is *not* what we're expecting as after 16
88+
# steps the override is shut off
89+
batched_shape = [1,1]
90+
incoming_grad_t = tf.constant(incoming_grad, shape=batched_shape)
91+
input_t = tf.constant(input, shape=batched_shape)
92+
with gradient_override_map({nonl_name: nonl_grad_override}):
93+
nonl_t = nonl(input_t)
94+
grad_wrt_input = tf.gradients(nonl_t, input_t, [incoming_grad_t])[0]
95+
nonl_t_no_override = nonl(input_t)
96+
grad_wrt_input_no_override = tf.gradients(nonl_t_no_override, input_t, [incoming_grad_t])[0]
97+
assert (grad_wrt_input.eval() == grad_wrt_input_no_override.eval()).all()

tests/optvis/test_integration.py

Lines changed: 5 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -20,4 +20,9 @@ def test_integration(decorrelate, fft):
2020
verbose=False, transforms=[])
2121
start_image = rendering[0]
2222
end_image = rendering[-1]
23+
objective_f = objectives.neuron("mixed3a", 177)
24+
param_f = lambda: param.image(64, decorrelate=decorrelate, fft=fft)
25+
rendering = render.render_vis(model, objective_f, param_f, verbose=False, thresholds=(0,64), use_fixed_seed=True)
26+
start_image, end_image = rendering
27+
2328
assert (start_image != end_image).any()

0 commit comments

Comments
 (0)