Skip to content

pnnl/neuromancer

Repository files navigation

NeuroMANCER v1.5.0

Neural Modules with Adaptive Nonlinear Constraints and Efficient Regularizations (NeuroMANCER) is an open-source differentiable programming (DP) library for solving parametric constrained optimization problems, physics-informed system identification, and parametric model-based optimal control. NeuroMANCER is written in PyTorch and allows for systematic integration of machine learning with scientific computing for creating end-to-end differentiable models and algorithms embedded with prior knowledge and physics.

⭐ Now available on PyPi! ⭐

Static Badge PyPI - Version

New in v1.5.0

Lightning Now supports integration with PyTorch Lightning (https://lightning.ai/docs/pytorch/stable/), bringing:

  • User workflow simplifications: zero boilerplate code and increased modularity
  • Ability for user to define custom training logic easily
  • Easy support for distributed GPU training
  • Weights and Biases hyperparameter tuning

Please refer to the Lightning folder and its README.

New Colab Examples:

Various domain examples, such as system identification of building thermal dynamics, in NeuroMANCER

PyTorch lightning integration Examples

Features and Examples

Extensive set of tutorials can be found in the examples folder. Interactive notebook versions of examples are available on Google Colab! Test out NeuroMANCER functionality before cloning the repository and setting up an environment.

Intro to NeuroMANCER

  • Open In Colab Part 1: Linear regression in PyTorch vs NeuroMANCER.

  • Open In Colab Part 2: NeuroMANCER syntax tutorial: variables, constraints, and objectives.

  • Open In Colab Part 3: NeuroMANCER syntax tutorial: modules, Node, and System class.

Parametric Programming

  • Open In Colab Part 1: Learning to solve a constrained optimization problem.

  • Open In Colab Part 2: Learning to solve a quadratically-constrained optimization problem.

  • Open In Colab Part 3: Learning to solve a set of 2D constrained optimization problems.

  • Open In Colab Part 4: Learning to solve a constrained optimization problem with the projected gradient.

  • Open In Colab Part 5: Using Cvxpylayers for differentiable projection onto the polytopic feasible set.

Ordinary Differential Equations (ODEs)

  • Open In Colab Part 1: Neural Ordinary Differential Equations (NODEs)
  • Open In Colab Part 2: Parameter estimation of ODE system
  • Open In Colab Part 3: Universal Differential Equations (UDEs)
  • Open In Colab Part 4: NODEs with exogenous inputs
  • Open In Colab Part 5: Neural State Space Models (NSSMs) with exogenous inputs
  • Open In Colab Part 6: Data-driven modeling of resistance-capacitance (RC) network ODEs
  • Open In Colab Part 7: Deep Koopman operator
  • Open In Colab Part 8: control-oriented Deep Koopman operator

Physics-Informed Neural Networks (PINNs) for Partial Differential Equations (PDEs)

  • Open In Colab Part 1: Diffusion Equation
  • Open In Colab Part 2: Burgers' Equation
  • Open In Colab Part 3: Burgers' Equation w/ Parameter Estimation (Inverse Problem)

Control

  • Open In Colab Part 1: Learning to stabilize a linear dynamical system.
  • Open In Colab Part 2: Learning to stabilize a nonlinear differential equation.
  • Open In Colab Part 3: Learning to control a nonlinear differential equation.
  • Open In Colab Part 4: Learning neural ODE model and control policy for an unknown dynamical system.
  • Open In Colab Part 5: Learning neural Lyapunov function for a nonlinear dynamical system.

Domain Examples

  • Open In Colab Part 1: Learning to Control Indoor Air Temperature in Buildings.
  • Open In Colab Part 2: Learning to Control an Pumped-Hydroelectricity Energy Storage System.
  • Open In Colab Part 3: Learning Building Thermal Dynamics using Neural ODEs.
  • Open In Colab Part 4: Data-driven modeling of a Resistance-Capacitance network with Neural ODEs.
  • Open In Colab Part 5: Learning Swing Equation Dynamics using Neural ODEs.
  • Open In Colab Part 6: Learning Building Thermal Dynamics using Neural State Space Models.

Lightning Integration Examples

  • Open In Colab Part 1: Lightning Integration Basics.
  • Open In Colab Part 2: Lightning Advanced Features and Automatic GPU Support.
  • Open In Colab Part 3: Hyperparameter Tuning With Lightning & WandB
  • Open In Colab Part 4: Defining Custom Training Logic via Lightning Modularized Code.

Documentation

The documentation for the library can be found online. There is also an introduction video covering core features of the library.

# Neuromancer syntax example for constrained optimization
import neuromancer as nm
import torch 

# define neural architecture 
func = nm.modules.blocks.MLP(insize=1, outsize=2, 
                             linear_map=nm.slim.maps['linear'], 
                             nonlin=torch.nn.ReLU, hsizes=[80] * 4)
# wrap neural net into symbolic representation via the Node class: map(p) -> x
map = nm.system.Node(func, ['p'], ['x'], name='map')
    
# define decision variables
x = nm.constraint.variable("x")[:, [0]]
y = nm.constraint.variable("x")[:, [1]]
# problem parameters sampled in the dataset
p = nm.constraint.variable('p')

# define objective function
f = (1-x)**2 + (y-x**2)**2
obj = f.minimize(weight=1.0)

# define constraints
con_1 = 100.*(x >= y)
con_2 = 100.*(x**2+y**2 <= p**2)

# create penalty method-based loss function
loss = nm.loss.PenaltyLoss(objectives=[obj], constraints=[con_1, con_2])
# construct differentiable constrained optimization problem
problem = nm.problem.Problem(nodes=[map], loss=loss)

UML diagram UML diagram of NeuroMANCER classes.

Installation

PIP Install (recommended)

Consider using a dedicated virtual environment (conda or otherwise) with Python 3.9+ installed.

pip install neuromancer

Example usage:

import torch
from neuromancer.system import Node

fun_1 = lambda x1, x2: 2.*x1 - x2**2
node_3 = Node(fun_1, ['y1', 'y2'], ['y3'], name='quadratic')
# evaluate forward pass of the node with dictionary input dataset
print(node_3({'y1': torch.rand(2), 'y2': torch.rand(2)}))

Manual Install

First clone the neuromancer package. A dedicated virtual environment (conda or otherwise) is recommended.

Note: If you have a previous neuromancer env it would be best at this point to create a new environment given the following instructions.

git clone -b master https://github.com/pnnl/neuromancer.git --single-branch

Create and activate virtual environment

conda create -n neuromancer python=3.10.4
conda activate neuromancer

Install neuromancer and all dependencies.

From top level directory of cloned neuromancer run:

pip install -e.[docs,tests,examples]

OR, for zsh users:

pip install -e.'[docs,tests,examples]'

See the pyproject.toml file for reference.

[project.optional-dependencies]
tests = ["pytest", "hypothesis"]
examples = ["casadi", "cvxpy", "imageio", "cvxpylayers"]
docs = ["sphinx", "sphinx-rtd-theme"]

Note on pip install with examples on MacOS (Apple M1)

Before CVXPY can be installed on Apple M1, you must install cmake via Homebrew:

brew install cmake

See CVXPY installation instructions for more details.

Conda install

Conda install is recommended for GPU acceleration.

❗️Warning: linux_env.yml, windows_env.yml, and osxarm64_env.yml are out of date. Manual installation of dependencies is recommended for conda.

Create environment & install dependencies

Ubuntu
conda env create -f linux_env.yml
conda activate neuromancer
Windows
conda env create -f windows_env.yml
conda activate neuromancer
conda install -c defaults intel-openmp -f
MacOS (Apple M1)
conda env create -f osxarm64_env.yml
conda activate neuromancer
Other (manually install all dependencies)

!!! Pay attention to comments for non-Linux OS !!!

conda create -n neuromancer python=3.10.4
conda activate neuromancer
conda install pytorch pytorch-cuda=11.6 -c pytorch -c nvidia
## OR (for Mac): conda install pytorch -c pytorch
conda config --append channels conda-forge
conda install scipy numpy"<1.24.0" matplotlib scikit-learn pandas dill mlflow pydot=1.4.2 pyts numba
conda install networkx=3.0 plum-dispatch=1.7.3 
conda install -c anaconda pytest hypothesis
conda install cvxpy cvxopt casadi seaborn imageio
conda install tqdm torchdiffeq toml
conda install lightning wandb -c conda-forge
## (for Windows): conda install -c defaults intel-openmp -f

Install NeuroMANCER package

From the top level directory of cloned neuromancer (in the activated environment where the dependencies have been installed):

pip install -e . --no-deps

Test NeuroMANCER install

Run pytest on the tests folder. It should take about 2 minutes to run the tests on CPU. There will be a lot of warnings that you can safely ignore. These warnings will be cleaned up in a future release.

Community Information

We welcome contributions and feedback from the open-source community!

Contributions, Discussions, and Issues

Please read the Community Development Guidelines for further information on contributions, discussions, and Issues.

Release notes

See the Release notes documenting new features.

License

NeuroMANCER comes with BSD license. See the license for further details.

Publications

Cite as

@article{Neuromancer2023,
  title={{NeuroMANCER: Neural Modules with Adaptive Nonlinear Constraints and Efficient Regularizations}},
  author={Drgona, Jan and Tuor, Aaron and Koch, James and Shapiro, Madelyn and Vrabie, Draguna},
  Url= {https://github.com/pnnl/neuromancer}, 
  year={2023}
}

Development team

Lead developers: Jan Drgona, Aaron Tuor
Active core developers: Madelyn Shapiro, James Koch, Rahul Birmiwal
Scientific advisors: Draguna Vrabie
Notable contributors: Seth Briney, Bo Tang, Ethan King, Shrirang Abhyankar, Mia Skomski, Stefan Dernbach, Zhao Chen, Christian Møldrup Legaard

Open-source contributions made by:

Made with contrib.rocks.

Acknowledgments

This research was partially supported by the Mathematics for Artificial Reasoning in Science (MARS) and Data Model Convergence (DMC) initiatives via the Laboratory Directed Research and Development (LDRD) investments at Pacific Northwest National Laboratory (PNNL), by the U.S. Department of Energy, through the Office of Advanced Scientific Computing Research's “Data-Driven Decision Control for Complex Systems (DnC2S)” project, and through the Energy Efficiency and Renewable Energy, Building Technologies Office under the “Dynamic decarbonization through autonomous physics-centric deep learning and optimization of building operations” and the “Advancing Market-Ready Building Energy Management by Cost-Effective Differentiable Predictive Control” projects. PNNL is a multi-program national laboratory operated for the U.S. Department of Energy (DOE) by Battelle Memorial Institute under Contract No. DE-AC05-76RL0-1830.