-
Notifications
You must be signed in to change notification settings - Fork 302
/
Copy pathindex.html
1122 lines (984 loc) · 53.2 KB
/
index.html
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
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
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
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
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
391
392
393
394
395
396
397
398
399
400
401
402
403
404
405
406
407
408
409
410
411
412
413
414
415
416
417
418
419
420
421
422
423
424
425
426
427
428
429
430
431
432
433
434
435
436
437
438
439
440
441
442
443
444
445
446
447
448
449
450
451
452
453
454
455
456
457
458
459
460
461
462
463
464
465
466
467
468
469
470
471
472
473
474
475
476
477
478
479
480
481
482
483
484
485
486
487
488
489
490
491
492
493
494
495
496
497
498
499
500
501
502
503
504
505
506
507
508
509
510
511
512
513
514
515
516
517
518
519
520
521
522
523
524
525
526
527
528
529
530
531
532
533
534
535
536
537
538
539
540
541
542
543
544
545
546
547
548
549
550
551
552
553
554
555
556
557
558
559
560
561
562
563
564
565
566
567
568
569
570
571
572
573
574
575
576
577
578
579
580
581
582
583
584
585
586
587
588
589
590
591
592
593
594
595
596
597
598
599
600
601
602
603
604
605
606
607
608
609
610
611
612
613
614
615
616
617
618
619
620
621
622
623
624
625
626
627
628
629
630
631
632
633
634
635
636
637
638
639
640
641
642
643
644
645
646
647
648
649
650
651
652
653
654
655
656
657
658
659
660
661
662
663
664
665
666
667
668
669
670
671
672
673
674
675
676
677
678
679
680
681
682
683
684
685
686
687
688
689
690
691
692
693
694
695
696
697
698
699
700
701
702
703
704
705
706
707
708
709
710
711
712
713
714
715
716
717
718
719
720
721
722
723
724
725
726
727
728
729
730
731
732
733
734
735
736
737
738
739
740
741
742
743
744
745
746
747
748
749
750
751
752
753
754
755
756
757
758
759
760
761
762
763
764
765
766
767
768
769
770
771
772
773
774
775
776
777
778
779
780
781
782
783
784
785
786
787
788
789
790
791
792
793
794
795
796
797
798
799
800
801
802
803
804
805
806
807
808
809
810
811
812
813
814
815
816
817
818
819
820
821
822
823
824
825
826
827
828
829
830
831
832
833
834
835
836
837
838
839
840
841
842
843
844
845
846
847
848
849
850
851
852
853
854
855
856
857
858
859
860
861
862
863
864
865
866
867
868
869
870
871
872
873
874
875
876
877
878
879
880
881
882
883
884
885
886
887
888
889
890
891
892
893
894
895
896
897
898
899
900
901
902
903
904
905
906
907
908
909
910
911
912
913
914
915
916
917
918
919
920
921
922
923
924
925
926
927
928
929
930
931
932
933
934
935
936
937
938
939
940
941
942
943
944
945
946
947
948
949
950
951
952
953
954
955
956
957
958
959
960
961
962
963
964
965
966
967
968
969
970
971
972
973
974
975
976
977
978
979
980
981
982
983
984
985
986
987
988
989
990
991
992
993
994
995
996
997
998
999
1000
<!DOCTYPE html>
<html lang="en">
<head>
<!-- Google Tag Manager -->
<script>(function(w,d,s,l,i){w[l]=w[l]||[];w[l].push({'gtm.start':
new Date().getTime(),event:'gtm.js'});var f=d.getElementsByTagName(s)[0],
j=d.createElement(s),dl=l!='dataLayer'?'&l='+l:'';j.async=true;j.src=
'https://www.googletagmanager.com/gtm.js?id='+i+dl;f.parentNode.insertBefore(j,f);
})(window,document,'script','dataLayer','GTM-T8XT4PS');</script>
<!-- End Google Tag Manager -->
<meta charset="UTF-8">
<meta name="viewport" content="width=device-width, initial-scale=1.0">
<meta http-equiv="X-UA-Compatible" content="ie=edge">
<link rel="shortcut icon" type="image/x-icon" href="/favicon.ico?">
<title>
Accelerated Generative Diffusion Models with PyTorch 2 | PyTorch
</title>
<meta name="robots" content="index, follow" />
<meta name="description" content="TL;DR: PyTorch 2.0 nightly offers out-of-the-box performance improvement for Generative Diffusion models by using the new torch.compile() compiler and optimized implementations of Multihead Attention integrated with PyTorch 2.
" />
<meta property="og:image" content="https://pytorch.org/assets/images/social-share.jpg" />
<meta name="twitter:image" content="https://pytorch.org/assets/images/social-share.jpg" />
<meta property="og:locale" content="en_US" />
<meta property="og:type" content="website" />
<meta property="og:title" content="Accelerated Generative Diffusion Models with PyTorch 2" />
<meta property="og:description" content="TL;DR: PyTorch 2.0 nightly offers out-of-the-box performance improvement for Generative Diffusion models by using the new torch.compile() compiler and optimized implementations of Multihead Attention integrated with PyTorch 2.
" />
<meta property="og:site_name" content="PyTorch" />
<meta name="twitter:card" content="summary_large_image" />
<meta name="twitter:title" content="Accelerated Generative Diffusion Models with PyTorch 2" />
<meta name="twitter:description" content="TL;DR: PyTorch 2.0 nightly offers out-of-the-box performance improvement for Generative Diffusion models by using the new torch.compile() compiler and optimized implementations of Multihead Attention integrated with PyTorch 2.
" />
<link rel="stylesheet" href="/assets/main.css">
<script src="/assets/vendor/jquery.min.js"></script>
<script src="/assets/vendor/popper.min.js"></script>
<script src="/assets/vendor/bootstrap.min.js"></script>
<script src="/assets/vendor/anchor.min.js"></script>
<script type="text/x-mathjax-config">
MathJax.Hub.Config({
tex2jax: {
skipTags: ['script', 'noscript', 'style', 'textarea', 'pre'],
inlineMath: [['$','$']]
}
});
</script>
<script src="https://cdn.mathjax.org/mathjax/latest/MathJax.js?config=TeX-AMS-MML_HTMLorMML" type="text/javascript"></script>
<script type="text/javascript" src="https://cdn.jsdelivr.net/npm/docsearch.js@2/dist/cdn/docsearch.min.js"></script>
<script>
!function(f,b,e,v,n,t,s)
{if(f.fbq)return;n=f.fbq=function(){n.callMethod?
n.callMethod.apply(n,arguments):n.queue.push(arguments)};
if(!f._fbq)f._fbq=n;n.push=n;n.loaded=!0;n.version='2.0';
n.queue=[];t=b.createElement(e);t.async=!0;
t.src=v;s=b.getElementsByTagName(e)[0];
s.parentNode.insertBefore(t,s)}(window,document,'script',
'https://connect.facebook.net/en_US/fbevents.js');
fbq('init', '243028289693773');
fbq('track', 'PageView');
</script>
<noscript>
<img height="1" width="1"
src="https://www.facebook.com/tr?id=243028289693773&ev=PageView
&noscript=1"/>
</noscript>
<!-- Twitter universal website tag code -->
<img height="1" width="1" style="display:none;" alt="" src="https://analytics.twitter.com/i/adsct?p_id=Twitter&p_user_id=0&txn_id=o2gi1&events=%5B%5B%22pageview%22%2Cnull%5D%5D&tw_sale_amount=0&tw_order_quantity=0 (https://urldefense.proofpoint.com/v2/url?u=https-3A__analytics.twitter.com_i_adsct-3Fp-5Fid-3DTwitter-26p-5Fuser-5Fid-3D0-26txn-5Fid-3Do2gi1-26events-3D-255B-255B-2522pageview-2522-252Cnull-255D-255D-26tw-5Fsale-5Famount-3D0-26tw-5Forder-5Fquantity-3D0&d=DwMGaQ&c=5VD0RTtNlTh3ycd41b3MUw&r=GMr8XYCDyeQQZuD3noL91A&m=dAJyokk16UvYy-vMrGn_JwYiGfp_eEgo25B9iGDCG-A&s=o6i4D0V0088WH2RnzIoqiF-vj45PL-2sTrsxQ0SNO3A&e=)" />
<img height="1" width="1" style="display:none;" alt="" src="//t.co/i/adsct?p_id=Twitter&p_user_id=0&txn_id=o2gi1&events=%5B%5B%22pageview%22%2Cnull%5D%5D&tw_sale_amount=0&tw_order_quantity=0 (https://urldefense.proofpoint.com/v2/url?u=https-3A__linkprotect.cudasvc.com_url-3Fa-3Dhttp-253a-252f-252ft.co-252fi-252fadsct-253fp-5Fid-253dTwitter-2526p-5Fuser-5Fid-253d0-2526txn-5Fid-253do2gi1-2526events-253d-25255B-25255B-252522pageview-252522-25252Cnull-25255D-25255D-2526tw-5Fsale-5Famount-253d0-2526tw-5Forder-5Fquantity-253d0-26c-3DE-2C1-2CC33dLwIhtuEcl5FhdztSnUwsioeej5k-2DWy0RYREBAq51kGji32A2Cw94YU9vQBpY5tPN3AukEw3C-5F-2DlbtndnLoR7-5FA-5FLoH0Rr7zLtP1ykptN-26typo-3D1&d=DwMGaQ&c=5VD0RTtNlTh3ycd41b3MUw&r=GMr8XYCDyeQQZuD3noL91A&m=dAJyokk16UvYy-vMrGn_JwYiGfp_eEgo25B9iGDCG-A&s=Abgc3XBkhESv8XBYtLchdDZyISGsK6v_BB6cLMJGyCw&e=)" />
<!-- End Twitter universal website tag code -->
<link rel="stylesheet" href="https://cdn.jsdelivr.net/npm/docsearch.js@2/dist/cdn/docsearch.min.css" />
<link href="/feed.xml" type="application/atom+xml" rel="alternate" title="Pythorch Blog Posts" />
</head>
<body class="blog">
<!-- Google Tag Manager (noscript) -->
<noscript><iframe src="https://www.googletagmanager.com/ns.html?id=GTM-T8XT4PS"
height="0" width="0" style="display:none;visibility:hidden"></iframe></noscript>
<!-- End Google Tag Manager (noscript) -->
<div class="main-background blog-background blog-detail-background"></div>
<div class="hello-bar">
<div class="container">
Join us at PyTorch Conference in San Francisco, October 22-23. CFP open now! <a target="_blank" href="https://events.linuxfoundation.org/pytorch-conference/">Learn more</a>.
</div>
</div>
<div class="container-fluid header-holder blog-detail-header">
<div class="container">
<div class="header-container">
<a class="header-logo" href="https://pytorch.org" aria-label="PyTorch"></a>
<div class="main-menu">
<ul>
<li class="main-menu-item">
<div id="dropdownMenuButton" data-toggle="resources-dropdown" class="resources-dropdown">
<a class="with-down-arrow">
Learn
</a>
<div class="resources-dropdown-menu">
<a class="nav-dropdown-item" href="/get-started">
<span class=dropdown-title>Get Started</span>
<p>Run PyTorch locally or get started quickly with one of the supported cloud platforms</p>
</a>
<a class="nav-dropdown-item" href="https://pytorch.org/tutorials/">
<span class="dropdown-title">Tutorials</span>
<p>Whats new in PyTorch tutorials</p>
</a>
<a class="nav-dropdown-item" href="https://pytorch.org/tutorials/beginner/basics/intro.html">
<span class="dropdown-title">Learn the Basics</span>
<p>Familiarize yourself with PyTorch concepts and modules</p>
</a>
<a class="nav-dropdown-item" href="https://pytorch.org/tutorials/recipes/recipes_index.html">
<span class="dropdown-title">PyTorch Recipes</span>
<p>Bite-size, ready-to-deploy PyTorch code examples</p>
</a>
<a class="nav-dropdown-item" href="https://pytorch.org/tutorials/beginner/introyt.html">
<span class="dropdown-title">Intro to PyTorch - YouTube Series</span>
<p>Master PyTorch basics with our engaging YouTube tutorial series</p>
</a>
<a class="nav-dropdown-item" href="/new">
<span class="dropdown-title">New to PyTorch Foundation</span>
</a>
</div>
</div>
</li>
<li class="main-menu-item">
<div id="dropdownMenuButton" data-toggle="resources-dropdown" class="resources-dropdown">
<a class="with-down-arrow">
Ecosystem
</a>
<div class="resources-dropdown-menu">
<a class="nav-dropdown-item" href="https://landscape.pytorch.org/" target="_blank">
<span class="dropdown-title">Tools</span>
<p>Learn about the tools and frameworks in the PyTorch Ecosystem</p>
</a>
<a class="nav-dropdown-item" href="/join-ecosystem">
<span class="dropdown-title">Join the Ecosystem</span>
</a>
<a class="nav-dropdown-item" href="/#community-module">
<span class=dropdown-title>Community</span>
<p>Join the PyTorch developer community to contribute, learn, and get your questions answered.</p>
</a>
<a class="nav-dropdown-item" href="https://discuss.pytorch.org" target="_blank">
<span class=dropdown-title>Forums</span>
<p>A place to discuss PyTorch code, issues, install, research</p>
</a>
<a class="nav-dropdown-item" href="/resources">
<span class=dropdown-title>Developer Resources</span>
<p>Find resources and get questions answered</p>
</a>
<a class="nav-dropdown-item" href="/ecosystem/contributor-awards-2024">
<span class="dropdown-title">Contributor Awards - 2024</span>
<p>Award winners announced at this year's PyTorch Conference</p>
</a>
</div>
</div>
</li>
<li class="main-menu-item">
<div id="dropdownMenuButton" data-toggle="resources-dropdown" class="resources-dropdown">
<a class="with-down-arrow">
Edge
</a>
<div class="resources-dropdown-menu">
<a class="nav-dropdown-item" href="/edge">
<span class="dropdown-title">About PyTorch Edge</span>
<p>Build innovative and privacy-aware AI experiences for edge devices</p>
</a>
<a class="nav-dropdown-item" href="/executorch-overview">
<span class="dropdown-title">ExecuTorch</span>
<p>End-to-end solution for enabling on-device inference capabilities across mobile and edge devices</p>
</a>
<a class="nav-dropdown-item" target="_blank" href="https://pytorch.org/executorch/stable/index.html">
<span class="dropdown-title">ExecuTorch Documentation</span>
</a>
</div>
</div>
</li>
<li class="main-menu-item">
<div id="docsDropdownButton" data-toggle="resources-dropdown" class="resources-dropdown">
<a class="with-down-arrow">
Docs
</a>
<div class="resources-dropdown-menu">
<a class="nav-dropdown-item" href="https://pytorch.org/docs">
<span class="dropdown-title">PyTorch</span>
<p>Explore the documentation for comprehensive guidance on how to use PyTorch.</p>
</a>
<a class="nav-dropdown-item" href="/pytorch-domains">
<span class="dropdown-title">PyTorch Domains</span>
<p> Read the PyTorch Domains documentation to learn more about domain-specific libraries.</p>
</a>
</div>
</div>
</li>
<li class="main-menu-item">
<div id="dropdownMenuButton" data-toggle="resources-dropdown" class="resources-dropdown">
<a class="with-down-arrow">
Blog & News
</a>
<div class="resources-dropdown-menu">
<a class="nav-dropdown-item" href="/blog">
<span class="dropdown-title">PyTorch Blog</span>
<p>Catch up on the latest technical news and happenings</p>
</a>
<a class="nav-dropdown-item" href="/community-blog">
<span class="dropdown-title">Community Blog</span>
<p>Stories from the PyTorch ecosystem</p>
</a>
<a class="nav-dropdown-item" href="/videos">
<span class="dropdown-title">Videos</span>
<p>Learn about the latest PyTorch tutorials, new, and more </p>
</a>
<a class="nav-dropdown-item" href="/community-stories">
<span class="dropdown-title">Community Stories</span>
<p>Learn how our community solves real, everyday machine learning problems with PyTorch</p>
</a>
<a class="nav-dropdown-item" href="/events">
<span class=dropdown-title>Events</span>
<p>Find events, webinars, and podcasts</p>
</a>
<a class="nav-dropdown-item" href="/newsletter">
<span class=dropdown-title>Newsletter</span>
<p>Stay up-to-date with the latest updates</p>
</a>
</div>
</div>
</li>
<li class="main-menu-item">
<div id="resourcesDropdownButton" data-toggle="resources-dropdown" class="resources-dropdown">
<a class="with-down-arrow">
About
</a>
<div class="resources-dropdown-menu">
<a class="nav-dropdown-item" href="/foundation">
<span class=dropdown-title>PyTorch Foundation</span>
<p>Learn more about the PyTorch Foundation.</p>
</a>
<a class="nav-dropdown-item" href="/governing-board">
<span class=dropdown-title>Governing Board</span>
</a>
<a class="nav-dropdown-item" href="/credits">
<span class=dropdown-title>Cloud Credit Program</span>
</a>
<a class="nav-dropdown-item" href="/tac">
<span class=dropdown-title>Technical Advisory Council</span>
</a>
<a class="nav-dropdown-item" href="/staff">
<span class=dropdown-title>Staff</span>
</a>
<a class="nav-dropdown-item" href="/contact-us">
<span class=dropdown-title>Contact Us</span>
</a>
</div>
</div>
</li>
<li class="main-menu-item">
<a href="/join" data-cta="join">
Become a Member
</a>
</li>
<li class="main-menu-item" id="github-main-menu-link">
<a href="https://github.com/pytorch/pytorch" title="Go to PyTorch GitHub">
<div id="topnav-gh-icon"></div>
</a>
</li>
<li class="navSearchWrapper reactNavSearchWrapper" key="search">
<div class="search-border">
<div id="search-icon"></div>
<input
id="search-input"
type="text"
title="Search"
/>
<div id="close-search">X</div>
</div>
</li>
</ul>
</div>
<script src="/assets/main-menu-dropdown.js"></script>
<a class="main-menu-open-button" href="#" data-behavior="open-mobile-menu"></a>
</div>
</div>
</div>
<div class="jumbotron jumbotron-fluid blog-detail-jumbotron">
<div class="container blog-detail-container">
<p class="featured-post">April 14, 2023</p>
<h1>
<a class="blog-title">Accelerated Generative Diffusion Models with PyTorch 2</a>
</h1>
</div>
</div>
<div class="main-content-wrapper blog-detail-wrapper">
<div class="main-content blog-detail-content">
<div class="container">
<img src="/assets/images/logo-icon.svg" class="img-fluid author-icon">
<article class="pytorch-article">
<p class="author">
by
Grigory Sizov, Michael Gschwind, Hamid Shojanazeri, Driss Guessous, Daniel Haziza, Christian Puhrsch
</p>
<p><strong>TL;DR</strong>: PyTorch 2.0 nightly offers out-of-the-box performance improvement for Generative Diffusion models by using the new <code class="language-plaintext highlighter-rouge">torch.compile()</code> compiler and optimized implementations of Multihead Attention integrated with PyTorch 2.</p>
<h2 id="introduction">Introduction</h2>
<p>A large part of the recent progress in Generative AI came from denoising diffusion models, which allow producing high quality images and videos from text prompts. This family includes Imagen, DALLE, Latent Diffusion, and others. However, all models in this family share a common drawback: generation is rather slow, due to the iterative nature of the sampling process by which the images are produced. This makes it important to optimize the code running inside the sampling loop.</p>
<p>We took an open source implementation of a popular text-to-image diffusion model as a starting point and accelerated its generation using two optimizations available in PyTorch 2: compilation and fast attention implementation. Together with a few minor memory processing improvements in the code these optimizations give up to 49% inference speedup relative to the original implementation without <a href="https://github.com/facebookresearch/xformers">xFormers</a>, and 39% inference speedup relative to using the original code with xFormers (excluding the compilation time), depending on the GPU architecture and batch size. Importantly, the speedup comes without a need to install xFormers or any other extra dependencies.</p>
<p>The table below shows the improvement in runtime between the original implementation with xFormers installed and our optimized version with PyTorch-integrated memory efficient attention (originally developed for and released in the <a href="https://github.com/facebookresearch/xformers">xFormers</a> library) and PyTorch compilation. The compilation time is excluded.</p>
<p><strong>Runtime improvement in % compared to original+xFormers</strong></p>
<p>See the absolute runtime numbers in section “Benchmarking setup and results summary”</p>
<table class="table table-bordered">
<thead>
<tr>
<td scope="col"><strong>GPU</strong>
</td>
<td scope="col"><strong>Batch size 1</strong>
</td>
<td scope="col"><strong>Batch size 2</strong>
</td>
<td scope="col"><strong>Batch size 4</strong>
</td>
</tr>
</thead>
<tr>
<td><strong>P100 (no compilation)</strong>
</td>
<td>-3.8
</td>
<td>0.44
</td>
<td>5.47
</td>
</tr>
<tr>
<td><strong>T4</strong>
</td>
<td>2.12
</td>
<td>10.51
</td>
<td>14.2
</td>
</tr>
<tr>
<td><strong>A10</strong>
</td>
<td>-2.34
</td>
<td>8.99
</td>
<td>10.57
</td>
</tr>
<tr>
<td><strong>V100</strong>
</td>
<td>18.63
</td>
<td>6.39
</td>
<td>10.43
</td>
</tr>
<tr>
<td><strong>A100</strong>
</td>
<td>38.5
</td>
<td>20.33
</td>
<td>12.17
</td>
</tr>
</table>
<p>One can notice the following:</p>
<ul>
<li>The improvements are significant for powerful GPUs like A100 and V100. For those GPUs the improvement is most pronounced for batch size 1</li>
<li>For less powerful GPUs we observe smaller speedups (or in two cases slight regressions). The batch size trend is reversed here: improvement is larger for larger batches</li>
</ul>
<p>In the following sections we describe the applied optimizations and provide detailed benchmarking data, comparing the generation time with various optimization features on/off.</p>
<p>Specifically, we benchmark 5 configurations and the plots below compare their absolute performance for different GPUs and batch sizes. For definitions of these configurations see section “Benchmarking setup and results”.</p>
<p><img src="/assets/images/2023-04-11-accelerated-generative-diffusion-models1.png" alt="Benchmark of denoising diffusion text-to-image generation across GPU architectures, batch size 1" style="max-height:800px; width:100%" /></p>
<p><img src="/assets/images/2023-04-11-accelerated-generative-diffusion-models2.png" alt="Benchmark of denoising diffusion text-to-image generation across GPU architectures, batch size 2" style="max-height:800px; width:100%" /></p>
<p><img src="/assets/images/2023-04-11-accelerated-generative-diffusion-models3.png" alt="Benchmark of denoising diffusion text-to-image generation across GPU architectures, batch size 1" style="max-height:800px; width:100%" /></p>
<h2 id="optimizations">Optimizations</h2>
<p>Here we’ll go into more detail about the optimizations introduced into the model code. These optimizations rely on features of PyTorch 2.0 which has been released recently.</p>
<h3 id="optimized-attention">Optimized Attention</h3>
<p>One part of the code which we optimized is the scaled dot-product attention. Attention is known to be a heavy operation: naive implementation materializes the attention matrix, leading to time and memory complexity quadratic in sequence length. It is common for diffusion models to use attention (<code class="language-plaintext highlighter-rouge">CrossAttention</code>) as part of Transformer blocks in multiple parts of the U-Net. Since the U-Net runs at every sampling step, this becomes a critical point to optimize. Instead of custom attention implementation one can use <code class="language-plaintext highlighter-rouge">torch.nn.MultiheadAttention,</code> which in PyTorch 2 has optimized attention implementation is integrated into it. This optimization schematically boils down to the following pseudocode:</p>
<div class="language-plaintext highlighter-rouge"><div class="highlight"><pre class="highlight"><code>class CrossAttention(nn.Module):
def __init__(self, ...):
# Create matrices: Q, K, V, out_proj
...
def forward(self, x, context=None, mask=None):
# Compute out = SoftMax(Q*K/sqrt(d))V
# Return out_proj(out)
…
</code></pre></div></div>
<p>gets replaced with</p>
<div class="language-plaintext highlighter-rouge"><div class="highlight"><pre class="highlight"><code>class CrossAttention(nn.Module):
def __init__(self, ...):
self.mha = nn.MultiheadAttention(...)
def forward(self, x, context):
return self.mha(x, context, context)
</code></pre></div></div>
<p>The optimized implementation of attention was available already in PyTorch 1.13 (see <a href="https://pytorch.org/blog/a-better-transformer-for-fast-transformer-encoder-inference/">here</a>) and widely adopted (see e.g. <a href="https://medium.com/pytorch/bettertransformer-out-of-the-box-performance-for-huggingface-transformers-3fbe27d50ab2">HuggingFace transformers library example</a>). In particular, it integrates memory-efficient attention from the <a href="https://github.com/facebookresearch/xformers">xFormers</a> library and flash attention from <a href="https://arxiv.org/abs/2205.14135">https://arxiv.org/abs/2205.14135</a>. PyTorch 2.0 expands this to additional attention functions such as cross attention and custom kernels for further acceleration, making it applicable to diffusion models.</p>
<p>Flash attention is available on GPUs with compute capability SM 7.5 or SM 8.x - for example, on T4, A10, and A100, which are included in our benchmark (you can check compute capability of each NVIDIA GPU <a href="https://developer.nvidia.com/cuda-gpus#compute">here</a>). However, in our tests on A100 the memory efficient attention performed better than flash attention for the particular case of diffusion models, due to the small number of attention heads and small batch size. PyTorch understands this and in this case chooses memory efficient attention over flash attention when both are available (see the logic <a href="https://github.com/pytorch/pytorch/blob/d8e795ecd53670682bd3b2e5ff1f378402b147d5/aten/src/ATen/native/transformers/cuda/sdp_utils.h#L33-L71">here</a>). For full control over the attention backends (memory-efficient attention, flash attention, “vanilla math”, or any future ones), power users can enable and disable them manually with the help of the context manager <a href="https://pytorch.org/docs/master/backends.html#torch.backends.cuda.sdp_kernel">torch.backends.cuda.sdp_kernel</a>.</p>
<h3 id="compilation">Compilation</h3>
<p>Compilation is a <a href="https://pytorch.org/get-started/pytorch-2.0/#user-experience">new feature of PyTorch 2.0</a>, enabling significant speedups with a very simple user experience. To invoke the default behavior, simply wrap a PyTorch module or a function into <code class="language-plaintext highlighter-rouge">torch.compile</code>:</p>
<div class="language-plaintext highlighter-rouge"><div class="highlight"><pre class="highlight"><code>model = torch.compile(model)
</code></pre></div></div>
<p>PyTorch compiler then turns Python code into a set of instructions which can be executed efficiently without Python overhead. The compilation happens dynamically the first time the code is executed. With the default behavior, under the hood PyTorch utilized <a href="https://pytorch.org/docs/stable/torch.compiler">TorchDynamo</a> to compile the code and <a href="https://dev-discuss.pytorch.org/t/torchinductor-a-pytorch-native-compiler-with-define-by-run-ir-and-symbolic-shapes/747">TorchInductor</a> to further optimize it. See <a href="https://pytorch.org/tutorials/intermediate/torch_compile_tutorial.html">this tutorial</a> for more details.</p>
<p>Although the one-liner above is enough for compilation, certain modifications in the code can squeeze a larger speedup. In particular, one should avoid so-called graph breaks - places in the code which PyTorch can’t compile. As opposed to previous PyTorch compilation approaches (like TorchScript), PyTorch 2 compiler doesn’t break in this case. Instead it falls back on eager execution - so the code runs, but with reduced performance. We introduced a few minor changes to the model code to get rid of graph breaks. This included eliminating functions from libraries not supported by the compiler, such as <code class="language-plaintext highlighter-rouge">inspect.isfunction</code> and <code class="language-plaintext highlighter-rouge">einops.rearrange</code>. See this <a href="https://pytorch.org/docs/stable/torch.compiler_faq.html#identifying-the-cause-of-a-graph-break">doc</a> to learn more about graph breaks and how to eliminate them.</p>
<p>Theoretically, one can apply <code class="language-plaintext highlighter-rouge">torch.compile </code>on the whole diffusion sampling loop. However, in practice it is enough to just compile the U-Net. The reason is that <code class="language-plaintext highlighter-rouge">torch.compile</code> doesn’t yet have a loop analyzer and would recompile the code for each iteration of the sampling loop. Moreover, compiled sampler code is likely to generate graph breaks - so one would need to adjust it if one wants to get a good performance from the compiled version.</p>
<p>Note that compilation <a href="https://github.com/openai/triton/blob/b5d32896b1f89fc44a82f8df3bb010934c53f4f5/README.md?plain=1#L66-L68">requires GPU compute capability >= SM 7.0</a> to run in non-eager mode. This covers all GPUs in our benchmarks - T4, V100, A10, A100 - except for P100 (see the <a href="https://developer.nvidia.com/cuda-gpus#compute">full list</a>).</p>
<h3 id="other-optimizations">Other optimizations</h3>
<p>In addition, we have improved efficiency of GPU memory operations by eliminating some common pitfalls, e.g. creating a tensor on GPU directly rather than creating it on CPU and later moving to GPU. The places where such optimizations were necessary were determined by line-profiling and looking at CPU/GPU traces and <a href="https://github.com/brendangregg/FlameGraph">Flame Graphs</a>.</p>
<h2 id="benchmarking-setup-and-results-summary">Benchmarking setup and results summary</h2>
<p>We have two versions of code to compare: <em>original</em> and <em>optimized</em>. On top of this, several optimization features (xFormers, PyTorch memory efficient attention, compilation) can be turned on/off. Overall, as mentioned in the introduction, we will be benchmarking 5 configurations:</p>
<ul>
<li><em>Original code without xFormers</em></li>
<li><em>Original code with xFormers</em></li>
<li><em>Optimized code with vanilla math attention backend and no compilation</em></li>
<li><em>Optimized code with memory-efficient attention backend and no compilation</em></li>
<li><em>Optimized code with memory-efficient attention backend and compilation</em></li>
</ul>
<p>As the <em>original version</em> we took the version of the code which uses PyTorch 1.12 and a custom implementation of attention. The <em>optimized version</em> uses <code class="language-plaintext highlighter-rouge">nn.MultiheadAttention</code> in <code class="language-plaintext highlighter-rouge">CrossAttention</code> and PyTorch 2.0.0.dev20230111+cu117. It also has a few other minor optimizations in PyTorch-related code.</p>
<p>The table below shows runtime of each version of the code in seconds, and the percentage improvement compared to the _original with xFormers. _The compilation time is excluded.</p>
<p><strong>Runtimes for batch size 1. In parenthesis - relative improvement with respect to the “Original with xFormers” row</strong></p>
<table class="table table-bordered">
<thead>
<tr>
<td><strong>Configuration</strong>
</td>
<td><strong>P100</strong>
</td>
<td><strong>T4</strong>
</td>
<td><strong>A10</strong>
</td>
<td><strong>V100</strong>
</td>
<td><strong>A100</strong>
</td>
</tr>
</thead>
<tr>
<td><strong>Original without xFormers</strong>
</td>
<td>30.4s (-19.3%)
</td>
<td>29.8s (-77.3%)
</td>
<td>13.0s (-83.9%)
</td>
<td>10.9s (-33.1%)
</td>
<td>8.0s (-19.3%)
</td>
</tr>
<tr>
<td><strong>Original with xFormers</strong>
</td>
<td><strong>25.5s</strong> (0.0%)
</td>
<td>16.8s (0.0%)
</td>
<td><strong>7.1s</strong> (0.0%)
</td>
<td>8.2s (0.0%)
</td>
<td>6.7s (0.0%)
</td>
</tr>
<tr>
<td><strong>Optimized with vanilla math attention, no compilation</strong>
</td>
<td>27.3s (-7.0%)
</td>
<td>19.9s (-18.7%)
</td>
<td>13.2s (-87.2%)
</td>
<td>7.5s (8.7%)
</td>
<td>5.7s (15.1%)
</td>
</tr>
<tr>
<td><strong>Optimized with mem. efficient attention, no compilation</strong>
</td>
<td>26.5s (-3.8%)
</td>
<td>16.8s (0.2%)
</td>
<td><strong>7.1s</strong> (-0.8%)
</td>
<td>6.9s (16.0%)
</td>
<td>5.3s (20.6%)
</td>
</tr>
<tr>
<td><strong>Optimized with mem. efficient attention and compilation</strong>
</td>
<td>-
</td>
<td><strong>16.4s </strong>(2.1%)
</td>
<td>7.2s (-2.3%)
</td>
<td><strong>6.6s</strong> (18.6%)
</td>
<td><strong>4.1s</strong> (38.5%)
</td>
</tr>
</table>
<p><strong>Runtimes for batch size 2</strong></p>
<table class="table table-bordered">
<thead>
<tr>
<td><strong>Configuration</strong>
</td>
<td><strong>P100</strong>
</td>
<td><strong>T4</strong>
</td>
<td><strong>A10</strong>
</td>
<td><strong>V100</strong>
</td>
<td><strong>A100</strong>
</td>
</tr>
</thead>
<tr>
<td><strong>Original without xFormers</strong>
</td>
<td>58.0s (-21.6%)
</td>
<td>57.6s (-84.0%)
</td>
<td>24.4s (-95.2%)
</td>
<td>18.6s (-63.0%)
</td>
<td>12.0s (-50.6%)
</td>
</tr>
<tr>
<td><strong>Original with xFormers</strong>
</td>
<td>47.7s (0.0%)
</td>
<td>31.3s (0.0%)
</td>
<td>12.5s (0.0%)
</td>
<td>11.4s (0.0%)
</td>
<td>8.0s (0.0%)
</td>
</tr>
<tr>
<td><strong>Optimized with vanilla math attention, no compilation</strong>
</td>
<td>49.3s (-3.5%)
</td>
<td>37.9s (-21.0%)
</td>
<td>17.8s (-42.2%)
</td>
<td>12.7s (-10.7%)
</td>
<td>7.8s (1.8%)
</td>
</tr>
<tr>
<td><strong>Optimized with mem. efficient attention, no compilation</strong>
</td>
<td><strong>47.5s </strong>(0.4%)
</td>
<td>31.2s (0.5%)
</td>
<td>12.2s (2.6%)
</td>
<td>11.5s (-0.7%)
</td>
<td>7.0s (12.6%)
</td>
</tr>
<tr>
<td><strong>Optimized with mem. efficient attention and compilation</strong>
</td>
<td>-
</td>
<td><strong>28.0s</strong> (10.5%)
</td>
<td><strong>11.4s</strong> (9.0%)
</td>
<td><strong>10.7s </strong>(6.4%)
</td>
<td><strong>6.4s</strong> (20.3%)
</td>
</tr>
</table>
<p><strong>Runtimes for batch size 4</strong></p>
<table class="table table-bordered">
<thead>
<tr>
<td><strong>Configuration</strong>
</td>
<td><strong>P100</strong>
</td>
<td><strong>T4</strong>
</td>
<td><strong>A10</strong>
</td>
<td><strong>V100</strong>
</td>
<td><strong>A100</strong>
</td>
</tr>
</thead>
<tr>
<td><strong>Original without xFormers</strong>
</td>
<td>117.9s (-20.0%)
</td>
<td>112.4s (-81.8%)
</td>
<td>47.2s (-101.7%)
</td>
<td>35.8s (-71.9%)
</td>
<td>22.8s (-78.9%)
</td>
</tr>
<tr>
<td><strong>Original with xFormers</strong>
</td>
<td>98.3s (0.0%)
</td>
<td>61.8s (0.0%)
</td>
<td>23.4s (0.0%)
</td>
<td>20.8s (0.0%)
</td>
<td>12.7s (0.0%)
</td>
</tr>
<tr>
<td><strong>Optimized with vanilla math attention, no compilation</strong>
</td>
<td>101.1s (-2.9%)
</td>
<td>73.0s (-18.0%)
</td>
<td>28.3s (-21.0%)
</td>
<td>23.3s (-11.9%)
</td>
<td>14.5s (-13.9%)
</td>
</tr>
<tr>
<td><strong>Optimized with mem. efficient attention, no compilation</strong>
</td>
<td><strong>92.9s </strong>(5.5%)
</td>
<td>61.1s (1.2%)
</td>
<td>23.9s (-1.9%)
</td>
<td>20.8s (-0.1%)
</td>
<td>12.8s (-0.9%)
</td>
</tr>
<tr>
<td><strong>Optimized with mem. efficient attention and compilation</strong>
</td>
<td>-
</td>
<td><strong>53.1s </strong>(14.2%)
</td>
<td><strong>20.9s</strong> (10.6%)
</td>
<td><strong>18.6s</strong> (10.4%)
</td>
<td><strong>11.2s</strong> (12.2%)
</td>
</tr>
</table>
<p>To minimize fluctuations and external influence on the performance of the benchmarked code, we ran each version of the code one after another, and then repeated this sequence 10 times: A, B, C, D, E, A, B, … So the results of a typical run would look like the one in the picture below.. Note that one shouldn’t rely on comparison of absolute run times between different graphs, but comparison of run times_ inside_ one graph is pretty reliable, thanks to our benchmarking setup.</p>
<p><img src="/assets/images/2023-04-11-accelerated-generative-diffusion-models4.png" alt="Denoising diffusion model generation benchmarks" style="max-height:700px" /></p>
<p>Each run of text-to-image generation script produces several batches, the number of which is regulated by the CLI parameter <code class="language-plaintext highlighter-rouge">--n_iter</code>. In the benchmarks we used <code class="language-plaintext highlighter-rouge">n_iter = 2</code>, but introduced an additional “warm-up” iteration, which doesn’t contribute to the run time. This was necessary for the runs with compilation, because compilation happens the first time the code runs, and so the first iteration is much longer than all subsequent. To make comparison fair, we also introduced this additional “warm-up” iteration to all other runs.</p>
<p>The numbers in the table above are for number of iterations 2 (plus a “warm-up one”), prompt ”A photo”, seed 1, PLMS sampler, and autocast turned on.</p>
<p>Benchmarks were done using P100, V100, A100, A10 and T4 GPUs. The T4 benchmarks were done in Google Colab Pro. The A10 benchmarks were done on g5.4xlarge AWS instances with 1 GPU.</p>
<h2 id="conclusions-and-next-steps">Conclusions and next steps</h2>
<p>We have shown that new features of PyTorch 2 - compiler and optimized attention implementation - give performance improvements exceeding or comparable with what previously required installation of an external dependency (xFormers). PyTorch achieved this, in particular, by integrating memory efficient attention from xFormers into its codebase. This is a significant improvement for user experience, given that xFormers, being a state-of-the-art library, in many scenarios requires custom installation process and long builds.</p>
<p>There are a few natural directions in which this work can be continued:</p>
<ul>
<li>The optimizations we implemented and described here are only benchmarked for text-to-image inference so far. It would be interesting to see how they affect training performance. PyTorch compilation can be directly applied to training; enabling training with PyTorch optimized attention is on the roadmap</li>
<li>We intentionally minimized changes to the original model code. Further profiling and optimization can probably bring more improvements</li>
<li>At the moment compilation is applied only to the U-Net model inside the sampler. Since there is a lot happening outside of U-Net (e.g. operations directly in the sampling loop), it would be beneficial to compile the whole sampler. However, this would require analysis of the compilation process to avoid recompilation at every sampling step</li>
<li>Current code only applies compilation within the PLMS sampler, but it should be trivial to extend it to other samplers</li>
<li>Besides text-to-image generation, diffusion models are also applied to other tasks - image-to-image and inpainting. It would be interesting to measure how their performance improves from PyTorch 2 optimizations</li>
</ul>
<p>See if you can increase performance of open source diffusion models using the methods we described, and share the results!</p>
<h2 id="resources">Resources</h2>
<ul>
<li>PyTorch 2.0 overview, which has a lot of information on <code class="language-plaintext highlighter-rouge">torch.compile:</code> <a href="https://pytorch.org/get-started/pytorch-2.0/">https://pytorch.org/get-started/pytorch-2.0/</a></li>
<li>Tutorial on <code class="language-plaintext highlighter-rouge">torch.compile</code>: <a href="https://pytorch.org/tutorials/intermediate/torch_compile_tutorial.html">https://pytorch.org/tutorials/intermediate/torch_compile_tutorial.html</a></li>
<li>General compilation troubleshooting: <a href="https://pytorch.org/docs/stable/torch.compiler_troubleshooting.html">https://pytorch.org/docs/stable/torch.compiler_troubleshooting.html</a></li>
<li>Details on graph breaks: <a href="https://pytorch.org/docs/stable/torch.compiler_faq.html#identifying-the-cause-of-a-graph-break">https://pytorch.org/docs/stable/torch.compiler_faq.html#identifying-the-cause-of-a-graph-break</a></li>
<li>Details on guards: <a href="https://pytorch.org/docs/stable/torch.compiler_guards_overview.html">https://pytorch.org/docs/stable/torch.compiler_guards_overview.html</a></li>
<li>Video deep dive on TorchDynamo <a href="https://www.youtube.com/watch?v=egZB5Uxki0I">https://www.youtube.com/watch?v=egZB5Uxki0I</a></li>
<li>Tutorial on optimized attention in PyTorch 1.12: <a href="https://pytorch.org/tutorials/beginner/bettertransformer_tutorial.html">https://pytorch.org/tutorials/beginner/bettertransformer_tutorial.html</a></li>
</ul>
<h2 id="acknowledgements">Acknowledgements</h2>
<p>We would like to thank Geeta Chauhan, Natalia Gimelshein, Patrick Labatut, Bert Maher, Mark Saroufim, Michael Voznesensky and Francisco Massa for their valuable advice and early feedback on the text.</p>
<p>Special thanks to Yudong Tao initiating the work on using PyTorch native attention in diffusion models.</p>
</article>
</div>
</div>
</div>
<!--
-->
<div class="container-fluid docs-tutorials-resources">
<div class="container">
<div class="row">
<div class="col-md-4 text-center">
<h2>Docs</h2>
<p>Access comprehensive developer documentation for PyTorch</p>
<a class="with-right-arrow" href="/docs">View Docs</a>
</div>
<div class="col-md-4 text-center">
<h2>Tutorials</h2>
<p>Get in-depth tutorials for beginners and advanced developers</p>
<a class="with-right-arrow" href="https://pytorch.org/tutorials">View Tutorials</a>
</div>
<div class="col-md-4 text-center">
<h2>Resources</h2>
<p>Find development resources and get your questions answered</p>
<a class="with-right-arrow" href="/resources">View Resources</a>
</div>
</div>
</div>
</div>
<footer class="site-footer">
<div class="container footer-container">
<div class="newsletter" id="newsletter">
<p
class="newsletter__title is-style-max-width-800"><strong>Stay in touch</strong> for updates, event info, and the latest news</p>
<script charset="utf-8" type="text/javascript" src="//js.hsforms.net/forms/embed/v2.js"></script>
<script>
hbspt.forms.create({
region: "na1",
portalId: "8112310",
formId: "2fb2231c-000b-4ec5-88a0-1ab242549c9e"
});
</script>
<p
class="newsletter__privacy">By submitting this form, I consent to receive marketing emails from the LF and its projects regarding their events, training, research, developments, and related announcements. I understand that I can unsubscribe at any time using the links in the footers of the emails I receive. <a href="https://www.linuxfoundation.org/privacy/">Privacy Policy</a>.</p>
</div>
<div class="lf-grid">
<div class="footer-logo-wrapper">
<a href="https://pytorch.org" class="footer-logo">
<img src="/assets/images/logo-icon.svg" alt="PyTorch logo" width="40">
</a>
</div>
<ul class="social-links">
<li><a href="https://www.facebook.com/pytorch" target="_blank" title="PyTorch on Facebook">
<svg xmlns="http://www.w3.org/2000/svg" viewbox="-0.51 -0.26 26.45 26.45" aria-label="Facebook"><path fill="currentColor" d="M25.497 13.075c0-2.45-.698-4.848-2.011-6.911a12.765 12.765 0 0 0-5.398-4.73A12.671 12.671 0 0 0 11.008.38a12.705 12.705 0 0 0-6.529 2.95A12.827 12.827 0 0 0 .563 9.358a12.896 12.896 0 0 0-.07 7.201 12.831 12.831 0 0 0 3.801 6.103 12.709 12.709 0 0 0 6.471 3.078v-8.957H7.53v-3.708h3.235v-2.824c0-3.213 1.903-4.988 4.813-4.988.956.014 1.909.097 2.852.25V8.67h-1.607a1.83 1.83 0 0 0-1.518.497 1.854 1.854 0 0 0-.561 1.505v2.404h3.535l-.563 3.708h-2.97v8.957a12.725 12.725 0 0 0 7.697-4.337 12.87 12.87 0 0 0 3.054-8.328z"/></svg>
</a></li>
<li><a href="https://twitter.com/pytorch" target="_blank" title="PyTorch on X">
<svg xmlns="http://www.w3.org/2000/svg" viewbox="0 0 300 300" aria-label="X"><path fill="currentColor" d="M178.57 127.15 290.27 0h-26.46l-97.03 110.38L89.34 0H0l117.13 166.93L0 300.25h26.46l102.4-116.59 81.8 116.59h89.34M36.01 19.54H76.66l187.13 262.13h-40.66"/></svg>
</a></li>
<li><a href="https://www.youtube.com/pytorch" target="_blank" title="PyTorch on YouTube">
<svg xmlns="http://www.w3.org/2000/svg" viewbox="0.21 0.27 34.45 25.07" aria-label="YouTube"><path fill="currentColor" d="M33.729 6.084s-.327-2.33-1.317-3.356a4.691 4.691 0 0 0-3.32-1.432c-4.634-.34-11.589-.34-11.589-.34h-.014s-6.954 0-11.59.342a4.692 4.692 0 0 0-3.32 1.432c-.993 1.025-1.315 3.354-1.315 3.354a52.189 52.189 0 0 0-.331 5.473v2.566c.014 1.829.125 3.656.331 5.472 0 0 .322 2.33 1.316 3.36 1.26 1.345 2.916 1.3 3.653 1.445 2.65.26 11.263.34 11.263.34s6.96-.01 11.597-.353a4.691 4.691 0 0 0 3.32-1.432c.993-1.026 1.316-3.356 1.316-3.356.206-1.817.316-3.644.33-5.473v-2.57a52.26 52.26 0 0 0-.33-5.472zM14.076 17.232V7.729l8.951 4.768-8.95 4.735z"/></svg>
</a></li>
<li><a href="https://www.linkedin.com/company/pytorch" target="_blank" title="PyTorch on LinkedIn">
<svg xmlns="http://www.w3.org/2000/svg" viewbox="-10.23 -10.23 531.96 531.96" aria-label="LinkedIn"><rect width="512" height="512" rx="0" fill="currentColor"/><circle fill="#000" cx="142" cy="138" r="37"/><path stroke="#000" stroke-width="66" d="M244 194v198M142 194v198"/><path fill="#000" d="M276 282c0-20 13-40 36-40 24 0 33 18 33 45v105h66V279c0-61-32-89-76-89-34 0-51 19-59 32"/></svg>
</a></li>
<li><a href="https://join.slack.com/t/pytorch/shared_invite/zt-2j2la612p-miUinTTaxXczKOJw48poHA" target="_blank" title="PyTorch Slack">
<svg xmlns="http://www.w3.org/2000/svg" viewBox="0.16 -0.03 21.19 21.19" aria-label="Slack"><path fill="currentColor" d="M4.896 13.27a2.147 2.147 0 0 1-2.141 2.142A2.147 2.147 0 0 1 .613 13.27c0-1.178.963-2.141 2.142-2.141h2.141v2.141zm1.08 0c0-1.178.962-2.141 2.141-2.141s2.142.963 2.142 2.141v5.363a2.147 2.147 0 0 1-2.142 2.141 2.147 2.147 0 0 1-2.141-2.142V13.27zm2.141-8.6a2.147 2.147 0 0 1-2.141-2.14c0-1.18.962-2.142 2.141-2.142s2.142.963 2.142 2.141v2.142H8.117zm0 1.08c1.179 0 2.141.962 2.141 2.141a2.147 2.147 0 0 1-2.141 2.142H2.755A2.147 2.147 0 0 1 .613 7.89c0-1.179.963-2.141 2.142-2.141h5.362zm8.599 2.141c0-1.179.963-2.141 2.141-2.141 1.179 0 2.143.962 2.143 2.14a2.147 2.147 0 0 1-2.142 2.142h-2.141V7.89zm-1.08 0a2.147 2.147 0 0 1-2.141 2.142 2.147 2.147 0 0 1-2.141-2.142V2.53c0-1.178.962-2.141 2.141-2.141s2.142.963 2.142 2.141v5.362zm-2.141 8.6c1.179 0 2.142.962 2.142 2.14a2.147 2.147 0 0 1-2.142 2.142 2.147 2.147 0 0 1-2.141-2.141V16.49h2.141zm0-1.08a2.147 2.147 0 0 1-2.141-2.141c0-1.179.962-2.142 2.141-2.142h5.362c1.179 0 2.142.963 2.142 2.142a2.147 2.147 0 0 1-2.142 2.142h-5.362z"></path></svg>
</a></li>
<li><a href="/wechat" title="PyTorch on WeChat">
<svg xmlns="http://www.w3.org/2000/svg" viewBox="0.14 -0.17 38.02 33.02" aria-label="WeChat"><path fill="currentColor" d="M26.289 10.976a12.972 12.972 0 0 0-8.742 3.53 10.386 10.386 0 0 0-3.224 8.795c-1.326-.164-2.535-.345-3.75-.448a2.332 2.332 0 0 0-1.273.216c-1.18.666-2.311 1.418-3.652 2.255.246-1.112.405-2.087.687-3.024a1.15 1.15 0 0 0-.523-1.52C1.737 17.902.02 13.601 1.307 9.165c1.189-4.1 4.11-6.587 8.077-7.884A13.54 13.54 0 0 1 24.18 5.617a10.135 10.135 0 0 1 2.109 5.359zM10.668 9.594a1.564 1.564 0 0 0-2.095-1.472 1.52 1.52 0 0 0-.895 1.964 1.502 1.502 0 0 0 1.391.966 1.545 1.545 0 0 0 1.598-1.46v.002zm8.15-1.566a1.567 1.567 0 0 0-1.528 1.543 1.528 1.528 0 0 0 1.571 1.492 1.52 1.52 0 0 0 1.375-2.117 1.518 1.518 0 0 0-1.415-.919l-.003.001z"></path><path fill="currentColor" d="M33.914 32.137c-1.075-.478-2.062-1.196-3.11-1.306-1.049-.11-2.145.494-3.24.605a10.821 10.821 0 0 1-8.781-2.864c-4.682-4.33-4.013-10.97 1.403-14.518 4.811-3.154 11.874-2.102 15.268 2.273a8.671 8.671 0 0 1-1.002 12.095c-1.046.929-1.422 1.693-.751 2.917.102.257.174.525.213.798zM21.68 20.292a1.264 1.264 0 1 0 .01-2.528 1.264 1.264 0 0 0-.01 2.528zm7.887-2.526a1.266 1.266 0 0 0-1.256 1.21 1.247 1.247 0 1 0 1.256-1.21z"></path></svg>
</a></li>
</ul>
</div>
<div class="privacy-policy">
<div class="copyright">
<p>© Copyright The Linux Foundation. The PyTorch Foundation is a project of The Linux Foundation.
For web site terms of use, trademark policy and other policies applicable to The PyTorch Foundation please see
<a href="https://www.linuxfoundation.org/legal/policies/">Linux Foundation Policies</a>. The PyTorch Foundation supports the PyTorch open source
project, which has been established as PyTorch Project a Series of LF Projects, LLC. For policies applicable to the PyTorch Project a Series of LF Projects, LLC,
please see <a href="https://www.lfprojects.org/policies/">LF Projects, LLC Policies</a>. <a href="https://www.linuxfoundation.org/privacy">Privacy Policy</a> and <a href="https://www.linuxfoundation.org/terms">Terms of Use</a>.</p>
</div>
</div>
</div>
</footer>
<div class="mobile-main-menu">
<div class="container-fluid">
<div class="container">
<div class="mobile-main-menu-header-container">
<a class="header-logo" href="https://pytorch.org" aria-label="PyTorch"></a>
<a class="main-menu-close-button" href="#" data-behavior="close-mobile-menu"></a>
</div>
</div>
</div>
<div class="mobile-main-menu-links-container">
<div class="main-menu">
<ul>
<li class="navSearchWrapper reactNavSearchWrapper tabletSearchWrapper" key="search">
<div class="mobile-search-border">
<input
id="mobile-search-input"
type="text"
title="Search"
/>
<div id="mobile-search-icon"></div>
</div>
</li>
<li class="resources-mobile-menu-title">
<a>Learn</a>
</li>
<ul class="resources-mobile-menu-items">
<li>
<a href="/get-started">Get Started</a>
</li>
<li>
<a href="https://pytorch.org/tutorials">Tutorials</a>
</li>
<li>
<a href="https://pytorch.org/tutorials/beginner/basics/intro.html">Learn the Basics</a>
</li>
<li>
<a href="https://pytorch.org/tutorials/recipes/recipes_index.html">PyTorch Recipes</a>
</li>
<li>
<a href="https://pytorch.org/tutorials/beginner/introyt.html">Introduction to PyTorch - YouTube Series</a>
</li>
<li>
<a href="/new">New to PyTorch Foundation</a>
</li>
</ul>
<li class="resources-mobile-menu-title">
<a>Ecosystem</a>
</li>
<ul class="resources-mobile-menu-items">
<li>
<a href="https://landscape.pytorch.org/">Tools</a>
</li>
<li>
<a href="/join-ecosystem">Join the Ecosystem</a>
</li>
<li>
<a href="/#community-module">Community</a>
</li>
<li>
<a href="https://discuss.pytorch.org">Forums</a>
</li>
<li>
<a href="/resources">Developer Resources</a>
</li>
<li>
<a href="/ecosystem/contributor-awards-2024">Contributor Awards - 2024</a>
</li>
</ul>
<li class="resources-mobile-menu-title">
<a>Edge</a>
</li>
<ul class="resources-mobile-menu-items">
<li>
<a href="/edge">About PyTorch Edge</a>
</li>
<li>
<a href="/executorch-overview">ExecuTorch</a>
</li>
<li>
<a href="https://pytorch.org/executorch/stable/index.html">ExecuTorch Documentation</a>
</li>
</ul>
<li class="resources-mobile-menu-title">
<a>Docs</a>
</li>
<ul class="resources-mobile-menu-items">