{ "cells": [ { "cell_type": "markdown", "source": [ "# Jacobians, Hessians, hvp, vhp, and more: composing functorch transforms\n", "\n", "\n", " \"Open\n", "\n", "\n", "Computing jacobians or hessians are useful in a number of non-traditional\n", "deep learning models. It is difficult (or annoying) to compute these quantities\n", "efficiently using a standard autodiff system like PyTorch Autograd; functorch\n", "provides ways of computing various higher-order autodiff quantities efficiently." ], "metadata": { "id": "zPbR6-eP51fe" }, "id": "zPbR6-eP51fe" }, { "cell_type": "markdown", "source": [ "## Computing the Jacobian" ], "metadata": { "id": "3kDj8fhn52j3" }, "id": "3kDj8fhn52j3" }, { "cell_type": "code", "source": [ "import torch\n", "import torch.nn as nn\n", "import torch.nn.functional as F\n", "from functools import partial\n", "_ = torch.manual_seed(0)" ], "metadata": { "id": "w_IinyjzflUH" }, "execution_count": null, "outputs": [], "id": "w_IinyjzflUH" }, { "cell_type": "markdown", "source": [ "Let’s start with a function that we’d like to compute the jacobian of. This is a simple linear function with non-linear activation.\n", "\n" ], "metadata": { "id": "cibF_PEYflUH" }, "id": "cibF_PEYflUH" }, { "cell_type": "code", "source": [ "def predict(weight, bias, x):\n", " return F.linear(x, weight, bias).tanh()" ], "metadata": { "id": "qhcD9hWYflUH" }, "execution_count": null, "outputs": [], "id": "qhcD9hWYflUH" }, { "cell_type": "markdown", "source": [ "Let's add some dummy data: a weight, a bias, and a feature vector x.\n", "\n" ], "metadata": { "id": "G8tqQrO_flUH" }, "id": "G8tqQrO_flUH" }, { "cell_type": "code", "source": [ "D = 16\n", "weight = torch.randn(D, D)\n", "bias = torch.randn(D)\n", "x = torch.randn(D) # feature vector" ], "metadata": { "id": "FZ4uJfZGflUH" }, "execution_count": null, "outputs": [], "id": "FZ4uJfZGflUH" }, { "cell_type": "markdown", "source": [ "Let's think of `predict` as a function that maps the input `x` from $R^D -> R^D$.\n", "PyTorch Autograd computes vector-Jacobian products. In order to compute the full\n", "Jacobian of this $R^D -> R^D$ function, we would have to compute it row-by-row\n", "by using a different unit vector each time." ], "metadata": { "id": "uMAW-ArQflUH" }, "id": "uMAW-ArQflUH" }, { "cell_type": "code", "source": [ "def compute_jac(xp):\n", " jacobian_rows = [torch.autograd.grad(predict(weight, bias, xp), xp, vec)[0]\n", " for vec in unit_vectors]\n", " return torch.stack(jacobian_rows)" ], "metadata": { "id": "z-BJPtbpflUI" }, "execution_count": null, "outputs": [], "id": "z-BJPtbpflUI" }, { "cell_type": "code", "source": [ "xp = x.clone().requires_grad_()\n", "unit_vectors = torch.eye(D)\n", "\n", "jacobian = compute_jac(xp)\n", "\n", "print(jacobian.shape)\n", "print(jacobian[0]) # show first row" ], "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "outputId": "f1f1ec12-56ef-40f7-8c3c-cbad7bf86644", "id": "zuWGSXspflUI" }, "execution_count": null, "outputs": [ { "output_type": "stream", "name": "stdout", "text": [ "torch.Size([16, 16])\n", "tensor([-0.5956, -0.6096, -0.1326, -0.2295, 0.4490, 0.3661, -0.1672, -1.1190,\n", " 0.1705, -0.6683, 0.1851, 0.1630, 0.0634, 0.6547, 0.5908, -0.1308])\n" ] } ], "id": "zuWGSXspflUI" }, { "cell_type": "markdown", "source": [ "Instead of computing the jacobian row-by-row, we can use vmap to get rid of the for-loop and vectorize the computation. \n", "We can’t directly apply vmap to PyTorch Autograd; instead, functorch provides a vjp transform:\n", "\n" ], "metadata": { "id": "mxlEOUieflUI" }, "id": "mxlEOUieflUI" }, { "cell_type": "code", "source": [ "from functorch import vmap, vjp\n", "\n", "_, vjp_fn = vjp(partial(predict, weight, bias), x)\n", "\n", "ft_jacobian, = vmap(vjp_fn)(unit_vectors)\n", "\n", "# lets confirm both methods compute the same result\n", "assert torch.allclose(ft_jacobian, jacobian)" ], "metadata": { "id": "DeF6uy4WflUI" }, "execution_count": null, "outputs": [], "id": "DeF6uy4WflUI" }, { "cell_type": "markdown", "source": [ "In future tutorial a composition of reverse-mode AD and vmap will give us per-sample-gradients. \n", "In this tutorial, composing reverse-mode AD and vmap gives us Jacobian computation! \n", "Various compositions of vmap and autodiff transforms can give us different interesting quantities.\n", "\n", "functorch provides **jacrev** as a convenience function that performs the vmap-vjp composition to compute jacobians. **jacrev** accepts an argnums argument that says which argument we would like to compute Jacobians with respect to.\n", "\n" ], "metadata": { "id": "Hy4REmwDflUI" }, "id": "Hy4REmwDflUI" }, { "cell_type": "code", "source": [ "from functorch import jacrev\n", "\n", "ft_jacobian = jacrev(predict, argnums=2)(weight, bias, x)\n", "\n", "# confirm \n", "assert torch.allclose(ft_jacobian, jacobian)" ], "metadata": { "id": "Rt7i6_YlflUI" }, "execution_count": null, "outputs": [], "id": "Rt7i6_YlflUI" }, { "cell_type": "markdown", "source": [ "Let’s compare the performance of the two ways to compute the jacobian. The functorch version is much faster (and becomes even faster the more outputs there are). \n", "\n", "In general, we expect that vectorization via vmap can help eliminate overhead and give better utilization of your hardware.\n", "\n", "Vmap does this magic by pushing the outer loop down into the functions primitive operations in order to obtain better performance.\n", "\n", "\n" ], "metadata": { "id": "JYe2H1UcflUJ" }, "id": "JYe2H1UcflUJ" }, { "cell_type": "markdown", "source": [ "Let's make a quick function to evaluate performance and deal with microseconds and milliseconds measurements:" ], "metadata": { "id": "i_143LZwflUJ" }, "id": "i_143LZwflUJ" }, { "cell_type": "code", "source": [ "def get_perf(first, first_descriptor, second, second_descriptor):\n", " \"\"\" takes torch.benchmark objects and compares delta of second vs first. \"\"\"\n", " faster = second.times[0]\n", " slower = first.times[0]\n", " gain = (slower-faster)/slower\n", " if gain < 0: gain *=-1 \n", " final_gain = gain*100\n", " print(f\" Performance delta: {final_gain:.4f} percent improvement with {second_descriptor} \")" ], "metadata": { "id": "II7r6jBtflUJ" }, "execution_count": null, "outputs": [], "id": "II7r6jBtflUJ" }, { "cell_type": "markdown", "source": [ "And then run the performance comparison:" ], "metadata": { "id": "r4clPnPKflUJ" }, "id": "r4clPnPKflUJ" }, { "cell_type": "code", "source": [ "from torch.utils.benchmark import Timer\n", "\n", "without_vmap = Timer(stmt=\"compute_jac(xp)\", globals=globals())\n", "with_vmap = Timer(stmt=\"jacrev(predict, argnums=2)(weight, bias, x)\", globals=globals())\n", "\n", "no_vmap_timer = without_vmap.timeit(500)\n", "with_vmap_timer = with_vmap.timeit(500)\n", "\n", "print(no_vmap_timer)\n", "print(with_vmap_timer)" ], "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "outputId": "cbf77a19-aac9-428d-eba1-74d337c53e49", "id": "ZPtoxF6eflUJ" }, "execution_count": null, "outputs": [ { "output_type": "stream", "name": "stdout", "text": [ "\n", "compute_jac(xp)\n", " 2.25 ms\n", " 1 measurement, 500 runs , 1 thread\n", "\n", "jacrev(predict, argnums=2)(weight, bias, x)\n", " 884.34 us\n", " 1 measurement, 500 runs , 1 thread\n" ] } ], "id": "ZPtoxF6eflUJ" }, { "cell_type": "markdown", "source": [ "Lets do a relative performance comparison of the above with our get_perf function:" ], "metadata": { "id": "nGBBi4dZflUJ" }, "id": "nGBBi4dZflUJ" }, { "cell_type": "code", "source": [ "get_perf(no_vmap_timer, \"without vmap\", with_vmap_timer, \"vmap\");" ], "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "outputId": "85d0bc5f-34aa-4826-f953-6c637404490c", "id": "zqV2RzEXflUJ" }, "execution_count": null, "outputs": [ { "output_type": "stream", "name": "stdout", "text": [ " Performance delta: 60.7170 percent improvement with vmap \n" ] } ], "id": "zqV2RzEXflUJ" }, { "cell_type": "markdown", "source": [ "Furthemore, it’s pretty easy to flip the problem around and say we want to compute Jacobians of the parameters to our model (weight, bias) instead of the input." ], "metadata": { "id": "EQAB99EQflUJ" }, "id": "EQAB99EQflUJ" }, { "cell_type": "code", "source": [ "# note the change in input via argnums params of 0,1 to map to weight and bias\n", "ft_jac_weight, ft_jac_bias = jacrev(predict, argnums=(0, 1))(weight, bias, x)" ], "metadata": { "id": "8UZpC8DnflUK" }, "execution_count": null, "outputs": [], "id": "8UZpC8DnflUK" }, { "cell_type": "markdown", "source": [ "## reverse-mode Jacobian (jacrev) vs forward-mode Jacobian (jacfwd)\n" ], "metadata": { "id": "F3USYENIflUK" }, "id": "F3USYENIflUK" }, { "cell_type": "markdown", "source": [ "We offer two APIs to compute jacobians: **jacrev** and **jacfwd**: \n", "- jacrev uses reverse-mode AD. As you saw above it is a composition of our vjp and vmap transforms. \n", "- jacfwd uses forward-mode AD. It is implemented as a composition of our jvp and vmap transforms. \n", "\n", "jacfwd and jacrev can be substituted for each other but they have different performance characteristics.\n", "\n", "As a general rule of thumb, if you’re computing the jacobian of an $𝑅^N \\to R^M$ function, and there are many more outputs than inputs (i.e. $M > N$) then jacfwd is preferred, otherwise use jacrev. There are exceptions to this rule, but a non-rigorous argument for this follows:\n", "\n", "In reverse-mode AD, we are computing the jacobian row-by-row, while in forward-mode AD (which computes Jacobian-vector products), we are computing it column-by-column. The Jacobian matrix has M rows and N columns, so if it is taller or wider one way we may prefer the method that deals with fewer rows or columns.\n", "\n" ], "metadata": { "id": "V7B3vE8dflUK" }, "id": "V7B3vE8dflUK" }, { "cell_type": "code", "source": [ "from functorch import jacrev, jacfwd" ], "metadata": { "id": "k7Tok7m3flUK" }, "execution_count": null, "outputs": [], "id": "k7Tok7m3flUK" }, { "cell_type": "markdown", "source": [ "First, let's benchmark with more inputs than outputs:\n", "\n" ], "metadata": { "id": "YrV-gZAaflUL" }, "id": "YrV-gZAaflUL" }, { "cell_type": "code", "source": [ "Din = 32\n", "Dout = 2048\n", "weight = torch.randn(Dout, Din)\n", "\n", "bias = torch.randn(Dout)\n", "x = torch.randn(Din)\n", "\n", "# remember the general rule about taller vs wider...here we have a taller matrix:\n", "print(weight.shape)\n", "\n", "using_fwd = Timer(stmt=\"jacfwd(predict, argnums=2)(weight, bias, x)\", globals=globals())\n", "using_bwd = Timer(stmt=\"jacrev(predict, argnums=2)(weight, bias, x)\", globals=globals())\n", "\n", "jacfwd_timing = using_fwd.timeit(500)\n", "jacrev_timing = using_bwd.timeit(500)\n", "\n", "print(f'jacfwd time: {jacfwd_timing}')\n", "print(f'jacrev time: {jacrev_timing}')\n" ], "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "outputId": "dd882726-9723-47c0-a72f-3c7835a85aa1", "id": "m5j-4hSxflUL" }, "execution_count": null, "outputs": [ { "output_type": "stream", "name": "stdout", "text": [ "torch.Size([2048, 32])\n", "jacfwd time: \n", "jacfwd(predict, argnums=2)(weight, bias, x)\n", " 1.32 ms\n", " 1 measurement, 500 runs , 1 thread\n", "jacrev time: \n", "jacrev(predict, argnums=2)(weight, bias, x)\n", " 12.46 ms\n", " 1 measurement, 500 runs , 1 thread\n" ] } ], "id": "m5j-4hSxflUL" }, { "cell_type": "markdown", "source": [ "and then do a relative benchmark:" ], "metadata": { "id": "k_Sg-4tVflUL" }, "id": "k_Sg-4tVflUL" }, { "cell_type": "code", "source": [ "get_perf(jacfwd_timing, \"jacfwd\", jacrev_timing, \"jacrev\", );" ], "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "outputId": "3a6586a1-269d-46d8-d119-e24f6d46277f", "id": "_4T96zGjflUL" }, "execution_count": null, "outputs": [ { "output_type": "stream", "name": "stdout", "text": [ " Performance delta: 842.8274 percent improvement with jacrev \n" ] } ], "id": "_4T96zGjflUL" }, { "cell_type": "markdown", "source": [ "and now the reverse - more outputs (M) than inputs (N):" ], "metadata": { "id": "RCDPot1yflUL" }, "id": "RCDPot1yflUL" }, { "cell_type": "code", "source": [ "Din = 2048\n", "Dout = 32\n", "weight = torch.randn(Dout, Din)\n", "bias = torch.randn(Dout)\n", "x = torch.randn(Din)\n", "\n", "using_fwd = Timer(stmt=\"jacfwd(predict, argnums=2)(weight, bias, x)\", globals=globals())\n", "using_bwd = Timer(stmt=\"jacrev(predict, argnums=2)(weight, bias, x)\", globals=globals())\n", "\n", "jacfwd_timing = using_fwd.timeit(500)\n", "jacrev_timing = using_bwd.timeit(500)\n", "\n", "print(f'jacfwd time: {jacfwd_timing}')\n", "print(f'jacrev time: {jacrev_timing}')" ], "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "outputId": "913e9ccd-3d4f-472a-a749-19cee36d0a16", "id": "_DRFqzqZflUM" }, "execution_count": null, "outputs": [ { "output_type": "stream", "name": "stdout", "text": [ "jacfwd time: \n", "jacfwd(predict, argnums=2)(weight, bias, x)\n", " 7.99 ms\n", " 1 measurement, 500 runs , 1 thread\n", "jacrev time: \n", "jacrev(predict, argnums=2)(weight, bias, x)\n", " 1.09 ms\n", " 1 measurement, 500 runs , 1 thread\n" ] } ], "id": "_DRFqzqZflUM" }, { "cell_type": "markdown", "source": [ "and a relative perf comparison:" ], "metadata": { "id": "5SRbMCNsflUM" }, "id": "5SRbMCNsflUM" }, { "cell_type": "code", "source": [ "get_perf(jacrev_timing, \"jacrev\", jacfwd_timing, \"jacfwd\")" ], "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "outputId": "c282ce25-4f6e-44cd-aed7-60f6f5010e5b", "id": "uF_9GaoiflUM" }, "execution_count": null, "outputs": [ { "output_type": "stream", "name": "stdout", "text": [ " Performance delta: 635.2095 percent improvement with jacfwd \n" ] } ], "id": "uF_9GaoiflUM" }, { "cell_type": "markdown", "source": [ "## Hessian computation with functorch.hessian\n" ], "metadata": { "id": "J29FQaBQflUM" }, "id": "J29FQaBQflUM" }, { "cell_type": "markdown", "source": [ "We offer a convenience API to compute hessians: `functorch.hessian`. \n", "Hessians are the jacobian of the jacobian (or the partial derivative of the partial derivative, aka second order).\n", "\n", "This suggests that one can just compose functorch’s jacobian transforms to compute the Hessian. \n", "Indeed, under the hood, `hessian(f)` is simply `jacfwd(jacrev(f))`.\n", "\n" ], "metadata": { "id": "My4DPH97flUM" }, "id": "My4DPH97flUM" }, { "cell_type": "markdown", "source": [ "Note: to boost performance: depending on your model, you may also want to use `jacfwd(jacfwd(f))` or `jacrev(jacrev(f))` instead to compute hessians leveraging the rule of thumb above regarding wider vs taller matrices.\n", "\n" ], "metadata": { "id": "FJt038l5flUM" }, "id": "FJt038l5flUM" }, { "cell_type": "code", "source": [ "from functorch import hessian\n", "\n", "# lets reduce the size in order not to blow out colab. Hessians require significant memory:\n", "Din = 512\n", "Dout = 32\n", "weight = torch.randn(Dout, Din)\n", "bias = torch.randn(Dout)\n", "x = torch.randn(Din)\n", "\n", "hess_api = hessian(predict, argnums=2)(weight, bias, x)\n", "hess_fwdfwd = jacfwd(jacfwd(predict, argnums=2), argnums=2)(weight, bias, x)\n", "#hess_revrev = jacrev(jacrev(predict, argnums=2), argnums=2)(weight, bias, x)\n" ], "metadata": { "id": "jEqr2ywZflUM" }, "execution_count": null, "outputs": [], "id": "jEqr2ywZflUM" }, { "cell_type": "markdown", "source": [ "Let's verify we have the same result regardless of using hessian api or using jacfwd(jacfwd())" ], "metadata": { "id": "n9BHcICQflUN" }, "id": "n9BHcICQflUN" }, { "cell_type": "code", "source": [ "torch.allclose(hess_api, hess_fwdfwd)" ], "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "outputId": "e457e3bc-f085-4f90-966d-f98893b98ea8", "id": "eHiWRkjJflUN" }, "execution_count": null, "outputs": [ { "output_type": "execute_result", "data": { "text/plain": [ "True" ] }, "metadata": {}, "execution_count": 18 } ], "id": "eHiWRkjJflUN" }, { "cell_type": "markdown", "source": [ "## Batch Jacobian and Batch Hessian\n" ], "metadata": { "id": "Gjt1RO8HflUN" }, "id": "Gjt1RO8HflUN" }, { "cell_type": "markdown", "source": [ "In the above examples we’ve been operating with a single feature vector. In some cases you might want to take the Jacobian of a batch of outputs with respect to a batch of inputs. That is, given a batch of inputs of shape `(B, N)` and a function that goes from $R^N \\to R^M$, we would like a Jacobian of shape `(B, M, N)`. \n", "\n", "The easiest way to do this is to use vmap:" ], "metadata": { "id": "RjIzdoQNflUN" }, "id": "RjIzdoQNflUN" }, { "cell_type": "code", "source": [ "batch_size = 64\n", "Din = 31\n", "Dout = 33\n", "\n", "weight = torch.randn(Dout, Din)\n", "print(f\"weight shape = {weight.shape}\")\n", "\n", "bias = torch.randn(Dout)\n", "\n", "x = torch.randn(batch_size, Din)" ], "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "outputId": "561eb618-e00f-40d5-bd99-fa51ab82051f", "id": "B1eoEO4UflUN" }, "execution_count": null, "outputs": [ { "output_type": "stream", "name": "stdout", "text": [ "weight shape = torch.Size([33, 31])\n" ] } ], "id": "B1eoEO4UflUN" }, { "cell_type": "code", "source": [ "compute_batch_jacobian = vmap(jacrev(predict, argnums=2), in_dims=(None, None, 0))\n", "batch_jacobian0 = compute_batch_jacobian(weight, bias, x)" ], "metadata": { "id": "nZ_V02NhflUN" }, "execution_count": null, "outputs": [], "id": "nZ_V02NhflUN" }, { "cell_type": "markdown", "source": [ "If you have a function that goes from (B, N) -> (B, M) instead and are certain that each input produces an independent output, then it’s also sometimes possible to do this without using vmap by summing the outputs and then computing the Jacobian of that function:\n", "\n" ], "metadata": { "id": "_OLDiY3MflUN" }, "id": "_OLDiY3MflUN" }, { "cell_type": "code", "source": [ "def predict_with_output_summed(weight, bias, x):\n", " return predict(weight, bias, x).sum(0)\n", "\n", "batch_jacobian1 = jacrev(predict_with_output_summed, argnums=2)(weight, bias, x).movedim(1, 0)\n", "assert torch.allclose(batch_jacobian0, batch_jacobian1)" ], "metadata": { "id": "_QH4hD8PflUO" }, "execution_count": null, "outputs": [], "id": "_QH4hD8PflUO" }, { "cell_type": "markdown", "source": [ "If you instead have a function that goes from $𝑅^𝑁 \\to 𝑅^𝑀$ but inputs that are batched, you compose vmap with jacrev to compute batched jacobians:\n", "\n", "Finally, batch hessians can be computed similarly. It’s easiest to think about them by using vmap to batch over hessian computation, but in some cases the sum trick also works.\n", "\n" ], "metadata": { "id": "eUjw65cCflUO" }, "id": "eUjw65cCflUO" }, { "cell_type": "code", "source": [ "compute_batch_hessian = vmap(hessian(predict, argnums=2), in_dims=(None, None, 0))\n", "\n", "batch_hess = compute_batch_hessian(weight, bias, x)\n", "batch_hess.shape" ], "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "outputId": "f3135cfa-e9e5-4f18-8cb7-0655e8a37cb5", "id": "3vAyQjMsflUO" }, "execution_count": null, "outputs": [ { "output_type": "execute_result", "data": { "text/plain": [ "torch.Size([64, 33, 31, 31])" ] }, "metadata": {}, "execution_count": 22 } ], "id": "3vAyQjMsflUO" }, { "cell_type": "markdown", "source": [ "## Computing Hessian-vector products\n", "\n", "The naive way to compute a Hessian-vector product (hvp) is to materialize the full Hessian and perform a dot-product with a vector. We can do better: it turns out we don't need to materialize the full Hessian to do this. We'll go through two (of many) different strategies to compute Hessian-vector products:\n", "- composing reverse-mode AD with reverse-mode AD\n", "- composing reverse-mode AD with forward-mode AD\n", "\n", "Composing reverse-mode AD with forward-mode AD (as opposed to reverse-mode with reverse-mode) is generally the more memory efficient way to compute a hvp because forward-mode AD doesn't need to construct an Autograd graph and save intermediates for backward:" ], "metadata": { "id": "Wa8E48sQgpkb" }, "id": "Wa8E48sQgpkb" }, { "cell_type": "code", "source": [ "from functorch import jvp, grad, vjp\n", "\n", "def hvp(f, primals, tangents):\n", " return jvp(grad(f), primals, tangents)[1]" ], "metadata": { "id": "trw6WbAth6BM" }, "execution_count": null, "outputs": [], "id": "trw6WbAth6BM" }, { "cell_type": "markdown", "source": [ "Here's some sample usage." ], "metadata": { "id": "DQMpRo6nitfr" }, "id": "DQMpRo6nitfr" }, { "cell_type": "code", "source": [ "def f(x):\n", " return x.sin().sum()\n", "\n", "x = torch.randn(2048)\n", "tangent = torch.randn(2048)\n", "\n", "result = hvp(f, (x,), (tangent,))" ], "metadata": { "id": "sPwg8SOdiVAK" }, "execution_count": null, "outputs": [], "id": "sPwg8SOdiVAK" }, { "cell_type": "markdown", "source": [ "If PyTorch forward-AD does not have coverage for your operations, then we can instead compose reverse-mode AD with reverse-mode AD:" ], "metadata": { "id": "zGvUIcB0j1Ez" }, "id": "zGvUIcB0j1Ez" }, { "cell_type": "code", "source": [ "def hvp_revrev(f, primals, tangents):\n", " _, vjp_fn = vjp(grad(f), *primals)\n", " return vjp_fn(*tangents)" ], "metadata": { "id": "mdDFZdlekAOK" }, "execution_count": null, "outputs": [], "id": "mdDFZdlekAOK" }, { "cell_type": "code", "source": [ "result_hvp_revrev = hvp_revrev(f, (x,), (tangent,))\n", "assert torch.allclose(result, result_hvp_revrev[0])" ], "metadata": { "id": "_CuCk9X0lW7C" }, "execution_count": null, "outputs": [], "id": "_CuCk9X0lW7C" } ], "metadata": { "kernelspec": { "display_name": "Python 3", "language": "python", "name": "python3" }, "language_info": { "codemirror_mode": { "name": "ipython", "version": 3 }, "file_extension": ".py", "mimetype": "text/x-python", "name": "python", "nbconvert_exporter": "python", "pygments_lexer": "ipython3", "version": "3.8.3" }, "colab": { "name": "jacobians_hessians.ipynb", "provenance": [] } }, "nbformat": 4, "nbformat_minor": 5 }