torch-harmonics
Overview | Installation | More information | Getting started | Contributors | Cite us | References
Overview
torch-harmonics is a differentiable implementation of the Spherical Harmonic transform in PyTorch. It was originally implemented to enable Spherical Fourier Neural Operators (SFNO). It uses quadrature rules to compute the projection onto the associated Legendre polynomials and FFTs for the projection onto the harmonic basis. This algorithm tends to outperform others with better asymptotic scaling for most practical purposes.
torch-harmonics uses PyTorch primitives to implement these operations, making it fully differentiable. Moreover, the quadrature can be distributed onto multiple ranks making it spatially distributed.
torch-harmonics has been used to implement a variety of differentiable PDE solvers which generated the animations below. Moreover, it has enabled the development of Spherical Fourier Neural Operators (SFNOs) [1].
Installation
Download directyly from PyPI:
pip install torch-harmonics
Build in your environment using the Python package:
git clone [email protected]:NVIDIA/torch-harmonics.git
cd torch-harmonics
pip install -e .
Alternatively, use the Dockerfile to build your custom container after cloning:
git clone [email protected]:NVIDIA/torch-harmonics.git
cd torch-harmonics
docker build . -t torch_harmonics
docker run --gpus all -it --rm --ipc=host --ulimit memlock=-1 --ulimit stack=67108864 torch_harmonics
More about torch-harmonics
Spherical harmonics
The spherical harmonics are special functions defined on the two-dimensional sphere
where
Spherical harmonic transform
The spherical harmonic transform (SHT)
realizes the projection of a signal
where
The implementation of the SHT follows the algorithm as presented in [2]. A direct spherical harmonic transform can be accomplished by a Fourier transform
in longitude and a Legendre transform
in latitude.
Discrete Legendre transform
The second integral, which computed the projection onto the Legendre polynomials is realized with quadrature. On the Gaussian grid, we use Gaussian quadrature in the
is obtained with the substitution
Here,
Getting started
The main functionality of torch_harmonics
is provided in the form of torch.nn.Modules
for composability. A minimum example is given by:
import torch
import torch_harmonics as th
device = torch.device('cuda' if torch.cuda.is_available() else 'cpu')
nlat = 512
nlon = 2*nlat
batch_size = 32
signal = torch.randn(batch_size, nlat, nlon)
# transform data on an equiangular grid
sht = th.RealSHT(nlat, nlon, grid="equiangular").to(device)
coeffs = sht(signal)
To enable scalable model-parallelism, torch-harmonics
implements a distributed variant of the SHT located in torch_harmonics.distributed
.
Detailed usage of torch-harmonics, alongside helpful analysis provided in a series of notebooks:
- Getting started
- Quadrature
- Visualizing the spherical harmonics
- Spectral fitting vs. SHT
- Conditioning of the Gramian
- Solving the Helmholtz equation
- Solving the shallow water equations
- Training Spherical Fourier Neural Operators
Remarks on automatic mixed precision (AMP) support
Note that torch-harmonics uses Fourier transforms from torch.fft
which in turn uses kernels from the optimized cuFFT
library. This library supports fourier transforms of float32
and float64
(i.e. single
and double
precision) tensors for all input sizes. For float16
(i.e. half
precision) and bfloat16
inputs however, the dimensions which are transformed are restricted to powers of two. Since data is converted to one of these reduced precision floating point formats when torch.cuda.amp.autocast
is used, torch-harmonics will issue an error when the input shapes are not powers of two. For these cases, we recommend disabling autocast for the harmonics transform specifically:
import torch
import torch_harmonics as th
sht = th.RealSHT(512, 1024, grid="equiangular").cuda()
with torch.cuda.amp.autocast(enabled = True):
# do some AMP converted math here
x = some_math(x)
# convert tensor to float32
x = x.to(torch.float32)
# now disable autocast specifically for the transform,
# making sure that the tensors are not converted
# back to reduced precision internally
with torch.cuda.amp.autocast(enabled = False):
xt = sht(x)
# continue operating on the transformed tensor
xt = some_more_math(xt)
Depending on the problem, it might be beneficial to upcast data to float64
instead of float32
precision for numerical stability.
Contributors
Boris Bonev ([email protected]), Thorsten Kurth ([email protected]), Christian Hundt ([email protected]), Nikola Kovachki ([email protected]), Jean Kossaifi ([email protected])
Cite us
If you use torch-harmonics
in an academic paper, please cite [1]
@misc{bonev2023spherical,
title={Spherical Fourier Neural Operators: Learning Stable Dynamics on the Sphere},
author={Boris Bonev and Thorsten Kurth and Christian Hundt and Jaideep Pathak and Maximilian Baust and Karthik Kashinath and Anima Anandkumar},
year={2023},
eprint={2306.03838},
archivePrefix={arXiv},
primaryClass={cs.LG}
}
References
[1] Bonev B., Kurth T., Hundt C., Pathak, J., Baust M., Kashinath K., Anandkumar A.; Spherical Fourier Neural Operators: Learning Stable Dynamics on the Sphere; arXiv 2306.0383, 2023.
[2] Schaeffer N.; Efficient spherical harmonic transforms aimed at pseudospectral numerical simulations; G3: Geochemistry, Geophysics, Geosystems, 2013.
[3] Wang B., Wang L., Xie Z.; Accurate calculation of spherical and vector spherical harmonic expansions via spectral element grids; Adv Comput Math, 2018.