fp8_primer.ipynb 18.1 KB
Newer Older
Przemek Tredak's avatar
Przemek Tredak committed
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
{
 "cells": [
  {
   "cell_type": "markdown",
   "id": "7b3e6954",
   "metadata": {},
   "source": [
    "# Using FP8 with Transformer Engine\n",
    "\n",
    "H100 GPU introduced support for a new datatype, FP8 (8-bit floating point), enabling higher throughput of matrix multiplies and convolutions. In this example we will introduce the FP8 datatype and show how to use it with Transformer Engine.\n",
    "\n",
    "## Introduction to FP8\n",
    "\n",
    "### Structure\n",
    "\n",
    "The FP8 datatype supported by H100 is actually 2 distinct datatypes, useful in different parts of the training of neural networks:\n",
    "\n",
    "* E4M3 - it consists of 1 sign bit, 4 exponent bits and 3 bits of mantissa. It can store values up to +/-448 and `nan`.\n",
    "* E5M2 - it consists of 1 sign bit, 5 exponent bits and 2 bits of mantissa. It can store values up to +/-57344, +/- `inf` and `nan`. The tradeoff of the increased dynamic range is lower precision of the stored values.\n",
    "\n",
21
    "<figure align=\"center\" id=\"fig_1\">\n",
Przemek Tredak's avatar
Przemek Tredak committed
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
    "<img src=\"fp8_formats.png\" width=\"60%\">\n",
    "<figcaption> Figure 1: Structure of the floating point datatypes. All of the values shown (in FP16, BF16, FP8 E4M3 and FP8 E5M2) are the closest representations of value 0.3952.</figcaption>\n",
    "</figure>\n",
    "\n",
    "During training neural networks both of these types may be utilized. Typically forward activations and weights require more precision, so E4M3 datatype is best used during forward pass. In the backward pass, however, gradients flowing through the network typically are less susceptible to the loss of precision, but require higher dynamic range. Therefore they are best stored using E5M2 data format. H100 TensorCores provide support for any combination of these types as the inputs, enabling us to store each tensor using its preferred precision.\n",
    "\n",
    "### Mixed precision training - a quick introduction\n",
    "\n",
    "In order to understand how FP8 can be used for training Deep Learning models, it is useful to first remind ourselves how mixed precision works with other datatypes, especially FP16.\n",
    "\n",
    "Mixed precision recipe for FP16 training has 2 components: choosing which operations should be performed in FP16 and dynamic loss scaling.\n",
    "\n",
    "* Choosing the operations to be performed in FP16 precision requires analysis of the numerical behavior of the outputs with respect to inputs of the operation as well as the expected performance benefit. This enables marking operations like matrix multiplies, convolutions and normalization layers as safe, while leaving `norm` or `exp` operations as requiring high precision.\n",
    "* Dynamic loss scaling enables avoiding both over- and underflows of the gradients during training. Those may happen since, while the dynamic range of FP16 is enough to store the distribution of the gradient values, this distribution may be centered around values too high or too low for FP16 to handle. Scaling the loss shifts those distributions (without affecting numerics by using only powers of 2) into the range representable in FP16. \n",
    "\n",
    "<figure align=\"center\">\n",
    "<img src=\"loss_scaling.png\" width=\"50%\">\n",
    "<figcaption> Figure 2: Scaling the loss enables shifting the gradient distribution into the representable range of FP16 datatype. </figcaption>\n",
    "</figure>\n",
    "\n",
    "### Mixed precision training with FP8\n",
    "\n",
    "While the dynamic range provided by the FP8 types is sufficient to store any particular activation or gradient, it is not sufficient for all of them at the same time. This makes the single loss scaling factor strategy, which worked for FP16, infeasible for FP8 training and instead requires using distinct scaling factors for each FP8 tensor.\n",
    "\n",
    "There are multiple strategies for choosing a scaling factor that is appropriate for a given FP8 tensor:\n",
    "\n",
    "* just-in-time scaling. This strategy chooses the scaling factor based on the maximum of absolute values (amax) of the tensor being produced. In practice it is infeasible, as it requires multiple passes through data - the operator produces and writes out the output in higher precision, then the maximum absolute value of the output is found and applied to all values in order to obtain the final FP8 output. This results in a lot of overhead, severely diminishing gains from using FP8.\n",
    "* delayed scaling. This strategy chooses the scaling factor based on the maximums of absolute values seen in some number of previous iterations. This enables full performance of FP8 computation, but requires storing the history of maximums as additional parameters of the FP8 operators. \n",
    "\n",
    "<figure align=\"center\">\n",
    "<img src=\"delayed_scaling.png\" width=\"80%\">\n",
    "<figcaption> Figure 3: Delayed scaling strategy. The FP8 operator uses scaling factor obtained using the history of amaxes (maximums of absolute values) seen in some number of previous iterations and produces both the FP8 output and the current amax, which gets stored in the history.</figcaption>\n",
    "</figure>\n",
    "\n",
    "As one can see in Figure 3, delayed scaling strategy requires both storing the history of amaxes, but also choosing a recipe for converting that history into the scaling factor used in the next iteration."
   ]
  },
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
  {
   "cell_type": "markdown",
   "id": "f03b58ed-71e8-422a-95be-35c1cc60c4e2",
   "metadata": {},
   "source": [
    "## MXFP8 and block scaling\n",
    "\n",
    "NVIDIA Blackwell architecture introduced support for a new variant of the FP8 format: [MXFP8](https://www.opencompute.org/documents/ocp-microscaling-formats-mx-v1-0-spec-final-pdf). \n",
    "\n",
    "### MXFP8 vs FP8\n",
    "\n",
    "The main difference between \"regular\" FP8 and MXFP8 lies in the granularity of the scaling. In FP8, each tensor has a single FP32 scaling factor, so all values in the tensor need to \"fit\" within the dynamic range of the FP8 datatype. This requires using the less precise E5M2 format to represent some tensors in the network (like gradients).\n",
    "\n",
    "MXFP8 addresses this by assigning a different scaling factor to each block of 32 [consecutive](#handling-transposes) values. This allows all values to be represented with the E4M3 datatype.\n",
    "\n",
    "<figure align=\"center\" id=\"fig_4\">\n",
    "<img src=\"MXFP8_FP8_comparison_1.png\" width=\"100%\">\n",
    "<figcaption> Figure 4: MXFP8 uses multiple scaling factors for a single tensor. The picture shows only 4 values per block for simplicity, but real MXFP8 has 32 values per block.</figcaption>\n",
    "</figure>\n",
    "\n",
    "<figure align=\"center\" id=\"fig_5\">\n",
    "<img src=\"MXFP8_FP8_comparison_2.png\" width=\"100%\">\n",
    "<figcaption> Figure 5: Due to multiple scaling factors, tensor's dynamic range requirements are reduced and so E4M3 format can be used as far fewer elements get saturated to 0.</figcaption>\n",
    "</figure>\n",
    "\n",
    "The second difference is the datatype used to store the scaling factors. FP8 uses FP32 (E8M23) while MXFP8 uses an 8-bit representation of a power of 2 (E8M0).\n",
    "\n",
    "<figure align=\"center\" id=\"fig_6\">\n",
    "<img src=\"E8M0.png\" width=\"100%\">\n",
    "<figcaption> Figure 6: Structure of the E8M0 datatype used for storing scaling factors in MXFP8.</figcaption>\n",
    "</figure>\n",
    "\n",
    "### Handling transposes\n",
    "\n",
    "The forward and backward passes of linear layers involve multiple matrix multiplications with different reduction dimensions. Blackwell Tensor Cores require MXFP8 data to be \"consecutive\" over the reduction dimension, so MXFP8 training uses non-transposed and transposed MXFP8 tensors at different points. However, while transposing FP8 data is numerically trivial, transposing MXFP8 data requires requantization.\n",
    "\n",
    "To avoid loss of precision connected with this double quantization, Transformer Engine creates both regular and transposed copies of the tensor from the original high precision input.\n",
    "\n",
    "<figure align=\"center\" id=\"fig_7\">\n",
    "<img src=\"linear_mxfp8.png\" width=\"80%\">\n",
    "<figcaption> Figure 7: Linear layer in MXFP8. Calculating both forward and backward pass requires tensors quantized in both directions.</figcaption>\n",
    "</figure>"
   ]
  },
Przemek Tredak's avatar
Przemek Tredak committed
103
104
105
106
107
108
109
  {
   "cell_type": "markdown",
   "id": "cf5e0b0d",
   "metadata": {},
   "source": [
    "## Using FP8 with Transformer Engine\n",
    "\n",
110
    "Transformer Engine library provides tools enabling easy to use training with FP8 datatype using FP8 delayed scaling and MXFP8 strategies.\n",
Przemek Tredak's avatar
Przemek Tredak committed
111
112
113
    "\n",
    "### FP8 recipe\n",
    "\n",
114
115
    "The [DelayedScaling](../api/common.rst#transformer_engine.common.recipe.DelayedScaling) recipe from the `transformer_engine.common.recipe` module stores all of the required options for training with FP8 delayed scaling: length of the amax history to use for scaling factor computation, FP8 data format, etc.\n",
    "Similarly, [MXFP8BlockScaling](../api/common.rst#transformer_engine.common.recipe.MXFP8BlockScaling) from the same module may be used to enable MXFP8 training."
Przemek Tredak's avatar
Przemek Tredak committed
116
117
118
119
120
121
122
123
124
   ]
  },
  {
   "cell_type": "code",
   "execution_count": 1,
   "id": "0c8fd0ef",
   "metadata": {},
   "outputs": [],
   "source": [
125
    "from transformer_engine.common.recipe import Format, DelayedScaling, MXFP8BlockScaling\n",
Przemek Tredak's avatar
Przemek Tredak committed
126
127
    "\n",
    "fp8_format = Format.HYBRID  # E4M3 during forward pass, E5M2 during backward pass\n",
128
129
130
    "fp8_recipe = DelayedScaling(fp8_format=fp8_format, amax_history_len=16, amax_compute_algo=\"max\")\n",
    "mxfp8_format = Format.E4M3  # E4M3 used everywhere\n",
    "mxfp8_recipe = MXFP8BlockScaling(fp8_format=mxfp8_format)"
Przemek Tredak's avatar
Przemek Tredak committed
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
382
383
384
385
386
387
388
389
390
   ]
  },
  {
   "cell_type": "markdown",
   "id": "f9591eb5",
   "metadata": {},
   "source": [
    "This recipe is then used to configure the FP8 training."
   ]
  },
  {
   "cell_type": "markdown",
   "id": "734d3934",
   "metadata": {},
   "source": [
    "### FP8 autocasting\n",
    "\n",
    "Not every operation is safe to be performed using FP8. All of the modules provided by Transformer Engine library were designed to provide maximum performance benefit from FP8 datatype while maintaining accuracy. In order to enable FP8 operations, TE modules need to be wrapped inside the [fp8_autocast](../api/pytorch.rst#transformer_engine.pytorch.fp8_autocast) context manager."
   ]
  },
  {
   "cell_type": "code",
   "execution_count": 2,
   "id": "f8b1ff7f",
   "metadata": {},
   "outputs": [],
   "source": [
    "import transformer_engine.pytorch as te\n",
    "import torch\n",
    "\n",
    "torch.manual_seed(12345)\n",
    "\n",
    "my_linear = te.Linear(768, 768, bias=True)\n",
    "\n",
    "inp = torch.rand((1024, 768)).cuda()\n",
    "\n",
    "with te.fp8_autocast(enabled=True, fp8_recipe=fp8_recipe):\n",
    "    out_fp8 = my_linear(inp)"
   ]
  },
  {
   "cell_type": "markdown",
   "id": "e41161f1",
   "metadata": {},
   "source": [
    "The `fp8_autocast` context manager hides the complexity of handling FP8:\n",
    "\n",
    "- All FP8-safe operations have their inputs cast to FP8\n",
    "- Amax history is updated\n",
    "- New scaling factors are computed and ready for the next iteration\n",
    "\n",
    "<div class=\"alert alert-info\">\n",
    "\n",
    "<b>Note</b>\n",
    "\n",
    "Support for FP8 in the Linear layer of Transformer Engine is currently limited to tensors with shapes where both dimensions are divisible by 16. In terms of the input to the full Transformer network, this typically requires padding sequence length to be multiple of 16.\n",
    "\n",
    "</div>"
   ]
  },
  {
   "cell_type": "markdown",
   "id": "f7bb2de9",
   "metadata": {},
   "source": [
    "### Handling backward pass\n",
    "\n",
    "When a model is run inside the `fp8_autocast` region, especially in multi-GPU training, some communication is required in order to synchronize the scaling factors and amax history. In order to perform that communication without introducing much overhead, `fp8_autocast` context manager aggregates the tensors before performing the communication.\n",
    "\n",
    "Due to this aggregation the backward call needs to happen outside of the `fp8_autocast` context manager. It has no impact on the computation precision - the precision of the backward pass is determined by the precision of the forward pass."
   ]
  },
  {
   "cell_type": "code",
   "execution_count": 3,
   "id": "e012bc8d",
   "metadata": {},
   "outputs": [],
   "source": [
    "loss_fp8 = out_fp8.mean()\n",
    "\n",
    "loss_fp8.backward()  # This backward pass uses FP8, since out_fp8 was calculated inside fp8_autocast\n",
    "\n",
    "out_fp32 = my_linear(inp)\n",
    "loss_fp32 = out_fp32.mean()\n",
    "loss_fp32.backward()  # This backward pass does not use FP8, since out_fp32 was calculated outside fp8_autocast"
   ]
  },
  {
   "cell_type": "markdown",
   "id": "1a6723ca",
   "metadata": {},
   "source": [
    "### Precision\n",
    "\n",
    "If we compare the results of the FP32 and FP8 execution, we will see that they are relatively close, but different:"
   ]
  },
  {
   "cell_type": "code",
   "execution_count": 4,
   "id": "41e9a37b",
   "metadata": {},
   "outputs": [
    {
     "data": {
      "text/plain": [
       "tensor([[ 0.2276,  0.2627,  0.3001,  ...,  0.0346,  0.2211,  0.1188],\n",
       "        [-0.0963, -0.3725,  0.1717,  ...,  0.0901,  0.0522, -0.3472],\n",
       "        [ 0.4526,  0.3482,  0.5976,  ..., -0.0687, -0.0382,  0.1566],\n",
       "        ...,\n",
       "        [ 0.1698,  0.6061,  0.0385,  ..., -0.2875, -0.1152, -0.0260],\n",
       "        [ 0.0679,  0.2946,  0.2751,  ..., -0.2284,  0.0517, -0.1441],\n",
       "        [ 0.1865,  0.2353,  0.9172,  ...,  0.1085,  0.1135,  0.1438]],\n",
       "       device='cuda:0', grad_fn=<_LinearBackward>)"
      ]
     },
     "execution_count": 4,
     "metadata": {},
     "output_type": "execute_result"
    }
   ],
   "source": [
    "out_fp8"
   ]
  },
  {
   "cell_type": "code",
   "execution_count": 5,
   "id": "b328ae0e",
   "metadata": {},
   "outputs": [
    {
     "data": {
      "text/plain": [
       "tensor([[ 0.2373,  0.2674,  0.2980,  ...,  0.0233,  0.2498,  0.1131],\n",
       "        [-0.0767, -0.3778,  0.1862,  ...,  0.0858,  0.0676, -0.3369],\n",
       "        [ 0.4615,  0.3593,  0.5813,  ..., -0.0779, -0.0349,  0.1422],\n",
       "        ...,\n",
       "        [ 0.1914,  0.6038,  0.0382,  ..., -0.2847, -0.0991, -0.0423],\n",
       "        [ 0.0864,  0.2895,  0.2719,  ..., -0.2388,  0.0772, -0.1541],\n",
       "        [ 0.2019,  0.2275,  0.9027,  ...,  0.1022,  0.1300,  0.1444]],\n",
       "       device='cuda:0', grad_fn=<_LinearBackward>)"
      ]
     },
     "execution_count": 5,
     "metadata": {},
     "output_type": "execute_result"
    }
   ],
   "source": [
    "out_fp32"
   ]
  },
  {
   "cell_type": "markdown",
   "id": "a9413c0a",
   "metadata": {},
   "source": [
    "That happens because in the FP8 case both the input and weights are cast to FP8 before the computation. We can see this if instead of the original inputs we use the inputs representable in FP8 (using a function defined in [quickstart_utils.py](quickstart_utils.py)):"
   ]
  },
  {
   "cell_type": "code",
   "execution_count": 6,
   "id": "ea939581",
   "metadata": {},
   "outputs": [
    {
     "name": "stdout",
     "output_type": "stream",
     "text": [
      "tensor([[ 0.2276,  0.2629,  0.3000,  ...,  0.0346,  0.2211,  0.1188],\n",
      "        [-0.0963, -0.3724,  0.1717,  ...,  0.0901,  0.0522, -0.3470],\n",
      "        [ 0.4526,  0.3479,  0.5976,  ..., -0.0686, -0.0382,  0.1566],\n",
      "        ...,\n",
      "        [ 0.1698,  0.6062,  0.0385,  ..., -0.2876, -0.1152, -0.0260],\n",
      "        [ 0.0679,  0.2947,  0.2750,  ..., -0.2284,  0.0516, -0.1441],\n",
      "        [ 0.1865,  0.2353,  0.9170,  ...,  0.1085,  0.1135,  0.1438]],\n",
      "       device='cuda:0', grad_fn=<_LinearBackward>)\n"
     ]
    }
   ],
   "source": [
    "from quickstart_utils import cast_to_representable\n",
    "\n",
    "inp_representable = cast_to_representable(inp)\n",
    "my_linear.weight.data = cast_to_representable(my_linear.weight.data)\n",
    "\n",
    "out_fp32_representable = my_linear(inp_representable)\n",
    "\n",
    "print(out_fp32_representable)"
   ]
  },
  {
   "cell_type": "markdown",
   "id": "03e703bd",
   "metadata": {},
   "source": [
    "This time the difference is really small:"
   ]
  },
  {
   "cell_type": "code",
   "execution_count": 7,
   "id": "78f1c2eb",
   "metadata": {},
   "outputs": [
    {
     "data": {
      "text/plain": [
       "tensor([[ 4.9591e-05, -1.9073e-04,  9.5367e-05,  ..., -3.8147e-06,\n",
       "          4.1962e-05,  2.2888e-05],\n",
       "        [ 2.2888e-05, -3.4332e-05,  2.2888e-05,  ...,  2.6703e-05,\n",
       "          5.3406e-05, -1.4114e-04],\n",
       "        [-3.8147e-05,  2.6703e-04, -3.8147e-06,  ..., -5.7220e-05,\n",
       "          4.1962e-05, -1.9073e-05],\n",
       "        ...,\n",
       "        [ 1.1444e-05, -7.2479e-05, -3.8147e-06,  ...,  5.3406e-05,\n",
       "         -1.5259e-05,  2.2888e-05],\n",
       "        [ 4.9591e-05, -9.5367e-05,  6.8665e-05,  ..., -1.5259e-05,\n",
       "          7.6294e-05,  4.5776e-05],\n",
       "        [-1.5259e-05, -7.6294e-06,  1.8692e-04,  ..., -3.0518e-05,\n",
       "         -4.5776e-05,  7.6294e-06]], device='cuda:0', grad_fn=<SubBackward0>)"
      ]
     },
     "execution_count": 7,
     "metadata": {},
     "output_type": "execute_result"
    }
   ],
   "source": [
    "out_fp8 - out_fp32_representable"
   ]
  },
  {
   "cell_type": "markdown",
   "id": "63ff9b8c",
   "metadata": {},
   "source": [
    "The differences in result coming from FP8 execution do not matter during the training process, but it is good to understand them, e.g. during debugging the model."
   ]
  }
 ],
 "metadata": {
  "kernelspec": {
   "display_name": "Python 3 (ipykernel)",
   "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",
391
   "version": "3.12.3"
Przemek Tredak's avatar
Przemek Tredak committed
392
393
394
395
396
  }
 },
 "nbformat": 4,
 "nbformat_minor": 5
}