CoCalc Logo Icon
StoreFeaturesDocsShareSupportNewsAboutSign UpSign In
pytorch

CoCalc provides the best real-time collaborative environment for Jupyter Notebooks, LaTeX documents, and SageMath, scalable from individual users to large groups and classes!

GitHub Repository: pytorch/tutorials
Path: blob/main/intermediate_source/neural_tangent_kernels.py
Views: 494
1
# -*- coding: utf-8 -*-
2
"""
3
Neural Tangent Kernels
4
======================
5
6
The neural tangent kernel (NTK) is a kernel that describes
7
`how a neural network evolves during training <https://en.wikipedia.org/wiki/Neural_tangent_kernel>`_.
8
There has been a lot of research around it `in recent years <https://arxiv.org/abs/1806.07572>`_.
9
This tutorial, inspired by the implementation of `NTKs in JAX <https://github.com/google/neural-tangents>`_
10
(see `Fast Finite Width Neural Tangent Kernel <https://arxiv.org/abs/2206.08720>`_ for details),
11
demonstrates how to easily compute this quantity using ``torch.func``,
12
composable function transforms for PyTorch.
13
14
.. note::
15
16
This tutorial requires PyTorch 2.0.0 or later.
17
18
Setup
19
-----
20
21
First, some setup. Let's define a simple CNN that we wish to compute the NTK of.
22
"""
23
24
import torch
25
import torch.nn as nn
26
from torch.func import functional_call, vmap, vjp, jvp, jacrev
27
device = 'cuda' if torch.cuda.device_count() > 0 else 'cpu'
28
29
class CNN(nn.Module):
30
def __init__(self):
31
super(CNN, self).__init__()
32
self.conv1 = nn.Conv2d(3, 32, (3, 3))
33
self.conv2 = nn.Conv2d(32, 32, (3, 3))
34
self.conv3 = nn.Conv2d(32, 32, (3, 3))
35
self.fc = nn.Linear(21632, 10)
36
37
def forward(self, x):
38
x = self.conv1(x)
39
x = x.relu()
40
x = self.conv2(x)
41
x = x.relu()
42
x = self.conv3(x)
43
x = x.flatten(1)
44
x = self.fc(x)
45
return x
46
47
######################################################################
48
# And let's generate some random data
49
50
x_train = torch.randn(20, 3, 32, 32, device=device)
51
x_test = torch.randn(5, 3, 32, 32, device=device)
52
53
######################################################################
54
# Create a function version of the model
55
# --------------------------------------
56
#
57
# ``torch.func`` transforms operate on functions. In particular, to compute the NTK,
58
# we will need a function that accepts the parameters of the model and a single
59
# input (as opposed to a batch of inputs!) and returns a single output.
60
#
61
# We'll use ``torch.func.functional_call``, which allows us to call an ``nn.Module``
62
# using different parameters/buffers, to help accomplish the first step.
63
#
64
# Keep in mind that the model was originally written to accept a batch of input
65
# data points. In our CNN example, there are no inter-batch operations. That
66
# is, each data point in the batch is independent of other data points. With
67
# this assumption in mind, we can easily generate a function that evaluates the
68
# model on a single data point:
69
70
71
net = CNN().to(device)
72
73
# Detaching the parameters because we won't be calling Tensor.backward().
74
params = {k: v.detach() for k, v in net.named_parameters()}
75
76
def fnet_single(params, x):
77
return functional_call(net, params, (x.unsqueeze(0),)).squeeze(0)
78
79
######################################################################
80
# Compute the NTK: method 1 (Jacobian contraction)
81
# ------------------------------------------------
82
# We're ready to compute the empirical NTK. The empirical NTK for two data
83
# points :math:`x_1` and :math:`x_2` is defined as the matrix product between the Jacobian
84
# of the model evaluated at :math:`x_1` and the Jacobian of the model evaluated at
85
# :math:`x_2`:
86
#
87
# .. math::
88
#
89
# J_{net}(x_1) J_{net}^T(x_2)
90
#
91
# In the batched case where :math:`x_1` is a batch of data points and :math:`x_2` is a
92
# batch of data points, then we want the matrix product between the Jacobians
93
# of all combinations of data points from :math:`x_1` and :math:`x_2`.
94
#
95
# The first method consists of doing just that - computing the two Jacobians,
96
# and contracting them. Here's how to compute the NTK in the batched case:
97
98
def empirical_ntk_jacobian_contraction(fnet_single, params, x1, x2):
99
# Compute J(x1)
100
jac1 = vmap(jacrev(fnet_single), (None, 0))(params, x1)
101
jac1 = jac1.values()
102
jac1 = [j.flatten(2) for j in jac1]
103
104
# Compute J(x2)
105
jac2 = vmap(jacrev(fnet_single), (None, 0))(params, x2)
106
jac2 = jac2.values()
107
jac2 = [j.flatten(2) for j in jac2]
108
109
# Compute J(x1) @ J(x2).T
110
result = torch.stack([torch.einsum('Naf,Mbf->NMab', j1, j2) for j1, j2 in zip(jac1, jac2)])
111
result = result.sum(0)
112
return result
113
114
result = empirical_ntk_jacobian_contraction(fnet_single, params, x_train, x_test)
115
print(result.shape)
116
117
######################################################################
118
# In some cases, you may only want the diagonal or the trace of this quantity,
119
# especially if you know beforehand that the network architecture results in an
120
# NTK where the non-diagonal elements can be approximated by zero. It's easy to
121
# adjust the above function to do that:
122
123
def empirical_ntk_jacobian_contraction(fnet_single, params, x1, x2, compute='full'):
124
# Compute J(x1)
125
jac1 = vmap(jacrev(fnet_single), (None, 0))(params, x1)
126
jac1 = jac1.values()
127
jac1 = [j.flatten(2) for j in jac1]
128
129
# Compute J(x2)
130
jac2 = vmap(jacrev(fnet_single), (None, 0))(params, x2)
131
jac2 = jac2.values()
132
jac2 = [j.flatten(2) for j in jac2]
133
134
# Compute J(x1) @ J(x2).T
135
einsum_expr = None
136
if compute == 'full':
137
einsum_expr = 'Naf,Mbf->NMab'
138
elif compute == 'trace':
139
einsum_expr = 'Naf,Maf->NM'
140
elif compute == 'diagonal':
141
einsum_expr = 'Naf,Maf->NMa'
142
else:
143
assert False
144
145
result = torch.stack([torch.einsum(einsum_expr, j1, j2) for j1, j2 in zip(jac1, jac2)])
146
result = result.sum(0)
147
return result
148
149
result = empirical_ntk_jacobian_contraction(fnet_single, params, x_train, x_test, 'trace')
150
print(result.shape)
151
152
######################################################################
153
# The asymptotic time complexity of this method is :math:`N O [FP]` (time to
154
# compute the Jacobians) + :math:`N^2 O^2 P` (time to contract the Jacobians),
155
# where :math:`N` is the batch size of :math:`x_1` and :math:`x_2`, :math:`O`
156
# is the model's output size, :math:`P` is the total number of parameters, and
157
# :math:`[FP]` is the cost of a single forward pass through the model. See
158
# section 3.2 in
159
# `Fast Finite Width Neural Tangent Kernel <https://arxiv.org/abs/2206.08720>`_
160
# for details.
161
#
162
# Compute the NTK: method 2 (NTK-vector products)
163
# -----------------------------------------------
164
#
165
# The next method we will discuss is a way to compute the NTK using NTK-vector
166
# products.
167
#
168
# This method reformulates NTK as a stack of NTK-vector products applied to
169
# columns of an identity matrix :math:`I_O` of size :math:`O\times O`
170
# (where :math:`O` is the output size of the model):
171
#
172
# .. math::
173
#
174
# J_{net}(x_1) J_{net}^T(x_2) = J_{net}(x_1) J_{net}^T(x_2) I_{O} = \left[J_{net}(x_1) \left[J_{net}^T(x_2) e_o\right]\right]_{o=1}^{O},
175
#
176
# where :math:`e_o\in \mathbb{R}^O` are column vectors of the identity matrix
177
# :math:`I_O`.
178
#
179
# - Let :math:`\textrm{vjp}_o = J_{net}^T(x_2) e_o`. We can use
180
# a vector-Jacobian product to compute this.
181
# - Now, consider :math:`J_{net}(x_1) \textrm{vjp}_o`. This is a
182
# Jacobian-vector product!
183
# - Finally, we can run the above computation in parallel over all
184
# columns :math:`e_o` of :math:`I_O` using ``vmap``.
185
#
186
# This suggests that we can use a combination of reverse-mode AD (to compute
187
# the vector-Jacobian product) and forward-mode AD (to compute the
188
# Jacobian-vector product) to compute the NTK.
189
#
190
# Let's code that up:
191
192
def empirical_ntk_ntk_vps(func, params, x1, x2, compute='full'):
193
def get_ntk(x1, x2):
194
def func_x1(params):
195
return func(params, x1)
196
197
def func_x2(params):
198
return func(params, x2)
199
200
output, vjp_fn = vjp(func_x1, params)
201
202
def get_ntk_slice(vec):
203
# This computes ``vec @ J(x2).T``
204
# `vec` is some unit vector (a single slice of the Identity matrix)
205
vjps = vjp_fn(vec)
206
# This computes ``J(X1) @ vjps``
207
_, jvps = jvp(func_x2, (params,), vjps)
208
return jvps
209
210
# Here's our identity matrix
211
basis = torch.eye(output.numel(), dtype=output.dtype, device=output.device).view(output.numel(), -1)
212
return vmap(get_ntk_slice)(basis)
213
214
# ``get_ntk(x1, x2)`` computes the NTK for a single data point x1, x2
215
# Since the x1, x2 inputs to ``empirical_ntk_ntk_vps`` are batched,
216
# we actually wish to compute the NTK between every pair of data points
217
# between {x1} and {x2}. That's what the ``vmaps`` here do.
218
result = vmap(vmap(get_ntk, (None, 0)), (0, None))(x1, x2)
219
220
if compute == 'full':
221
return result
222
if compute == 'trace':
223
return torch.einsum('NMKK->NM', result)
224
if compute == 'diagonal':
225
return torch.einsum('NMKK->NMK', result)
226
227
# Disable TensorFloat-32 for convolutions on Ampere+ GPUs to sacrifice performance in favor of accuracy
228
with torch.backends.cudnn.flags(allow_tf32=False):
229
result_from_jacobian_contraction = empirical_ntk_jacobian_contraction(fnet_single, params, x_test, x_train)
230
result_from_ntk_vps = empirical_ntk_ntk_vps(fnet_single, params, x_test, x_train)
231
232
assert torch.allclose(result_from_jacobian_contraction, result_from_ntk_vps, atol=1e-5)
233
234
######################################################################
235
# Our code for ``empirical_ntk_ntk_vps`` looks like a direct translation from
236
# the math above! This showcases the power of function transforms: good luck
237
# trying to write an efficient version of the above by only using
238
# ``torch.autograd.grad``.
239
#
240
# The asymptotic time complexity of this method is :math:`N^2 O [FP]`, where
241
# :math:`N` is the batch size of :math:`x_1` and :math:`x_2`, :math:`O` is the
242
# model's output size, and :math:`[FP]` is the cost of a single forward pass
243
# through the model. Hence this method performs more forward passes through the
244
# network than method 1, Jacobian contraction (:math:`N^2 O` instead of
245
# :math:`N O`), but avoids the contraction cost altogether (no :math:`N^2 O^2 P`
246
# term, where :math:`P` is the total number of model's parameters). Therefore,
247
# this method is preferable when :math:`O P` is large relative to :math:`[FP]`,
248
# such as fully-connected (not convolutional) models with many outputs :math:`O`.
249
# Memory-wise, both methods should be comparable. See section 3.3 in
250
# `Fast Finite Width Neural Tangent Kernel <https://arxiv.org/abs/2206.08720>`_
251
# for details.
252
253