Skip to content

Commit

Permalink
Parametric DMD for heat conduction example (#6)
Browse files Browse the repository at this point in the history
* StopWatch python class added in python_utils module.

* PointwiseSnapshot python class.

* templetized pyDatabase.

* bindings for CSV/HDFDatabase. need to finish all member functions.

* binding of some member functions needed for parametric dmd.

* templatized getVectorPointer. putDoubleArray uses numpy array.

* ParametricDMD bindings.

* enforce noconvert() for DMD.train() arguments.

* templatized getParametricDMD.

* parametric dmd for heat conduction example.

* separated StopWatch class.

* updated example command lines.
  • Loading branch information
dreamer2368 authored Aug 24, 2023
1 parent 17052e9 commit b9581d4
Show file tree
Hide file tree
Showing 14 changed files with 1,498 additions and 199 deletions.
17 changes: 14 additions & 3 deletions CMakeLists.txt
Original file line number Diff line number Diff line change
Expand Up @@ -19,6 +19,9 @@ if (BUILD_SCALAPACK)
add_custom_target(RUN_SLPK_BUILD ALL DEPENDS SLPK_BUILD)
endif(BUILD_SCALAPACK)

# HDF5
find_package(HDF5 1.8.0 REQUIRED)

#=================== libROM ================================

#It is tedious to build libROM. option to not build
Expand Down Expand Up @@ -108,17 +111,19 @@ include_directories(
${LIBROM_INCLUDE_DIR}
${MPI_INCLUDE_PATH}
${MPI4PY}
${HDF5_C_INCLUDE_DIRS}
# ${MFEM_INCLUDES}
# ${HYPRE_INCLUDES}
# ${PARMETIS_INCLUDES}
# ${MFEM_C_INCLUDE_DIRS}
)
# link_libraries(
link_libraries(
${HDF5_LIBRARIES}
# ${MFEM}
# ${HYPRE}
# ${PARMETIS}
# ${METIS}
# )
)

add_subdirectory("extern/pybind11")

Expand All @@ -135,9 +140,15 @@ pybind11_add_module(_pylibROM
bindings/pylibROM/linalg/svd/pySVD.cpp
bindings/pylibROM/linalg/svd/pyStaticSVD.cpp
bindings/pylibROM/linalg/svd/pyIncrementalSVD.cpp
bindings/pylibROM/algo/pyDMD.cpp

bindings/pylibROM/algo/pyDMD.cpp
bindings/pylibROM/algo/pyParametricDMD.cpp

bindings/pylibROM/utils/mpi_utils.cpp
bindings/pylibROM/utils/pyDatabase.hpp
bindings/pylibROM/utils/pyDatabase.cpp
bindings/pylibROM/utils/pyHDFDatabase.cpp
bindings/pylibROM/utils/pyCSVDatabase.cpp

# TODO(kevin): We do not bind mfem-related functions until we figure out how to type-cast SWIG Object.
# Until then, mfem-related functions need to be re-implemented on python-end, using PyMFEM.
Expand Down
33 changes: 3 additions & 30 deletions bindings/pylibROM/algo/pyDMD.cpp
Original file line number Diff line number Diff line change
Expand Up @@ -7,40 +7,13 @@
#include <pybind11/stl.h>
#include "algo/DMD.h"
#include "linalg/Vector.h"
#include "linalg/Matrix.h"
#include "python_utils/cpp_utils.hpp"

namespace py = pybind11;
using namespace CAROM;
using namespace std;

/*
//wrapper class needed for setOffset and takeSample methods.
class PyDMD : public DMD
{
public:
using DMD::DMD;
void setOffset(py::object offset_vector, int order)
{
// Convert the Python object to the appropriate type
Vector* offset = py::cast<Vector*>(offset_vector);
DMD::setOffset(offset, order);
}
void takeSample(py::array_t<double> u_in, double t)
{
// Access the underlying data of the NumPy array
auto buf = u_in.request();
double* u_in_ptr = static_cast<double*>(buf.ptr);
// Call the original method
DMD::takeSample(u_in_ptr, t);
}
// Add wrappers for other methods as needed
};
*/

void init_DMD(pybind11::module_ &m) {

py::class_<DMD>(m, "DMD")
Expand All @@ -58,9 +31,9 @@ void init_DMD(pybind11::module_ &m) {
self.takeSample(getVectorPointer(u_in), t);
})
.def("train", py::overload_cast<double, const Matrix*, double>(&DMD::train),
py::arg("energy_fraction"), py::arg("W0") = nullptr, py::arg("linearity_tol") = 0.0)
py::arg("energy_fraction").noconvert(), py::arg("W0") = nullptr, py::arg("linearity_tol") = 0.0)
.def("train", py::overload_cast<int, const Matrix*, double>(&DMD::train),
py::arg("k"), py::arg("W0") = nullptr, py::arg("linearity_tol") = 0.0)
py::arg("k").noconvert(), py::arg("W0") = nullptr, py::arg("linearity_tol") = 0.0)
.def("projectInitialCondition", &DMD::projectInitialCondition,
py::arg("init"), py::arg("t_offset") = -1.0)
.def("predict", &DMD::predict, py::arg("t"), py::arg("deg") = 0)
Expand Down
123 changes: 123 additions & 0 deletions bindings/pylibROM/algo/pyParametricDMD.cpp
Original file line number Diff line number Diff line change
@@ -0,0 +1,123 @@
//
// Created by barrow9 on 6/4/23.
//
#include <pybind11/pybind11.h>
#include <pybind11/numpy.h>
#include <pybind11/operators.h>
#include <pybind11/stl.h>
#include "algo/DMD.h"
#include "algo/ParametricDMD.h"
#include "linalg/Vector.h"
#include "python_utils/cpp_utils.hpp"

namespace py = pybind11;
using namespace CAROM;
using namespace std;

template <class T>
T* getParametricDMD_Type(
std::vector<Vector*>& parameter_points,
std::vector<T*>& dmds,
Vector* desired_point,
std::string rbf,
std::string interp_method,
double closest_rbf_val,
bool reorthogonalize_W)
{
T *parametric_dmd = NULL;
getParametricDMD(parametric_dmd, parameter_points, dmds, desired_point,
rbf, interp_method, closest_rbf_val, reorthogonalize_W);
return parametric_dmd;
}

template <class T>
T* getParametricDMD_Type(
std::vector<Vector*>& parameter_points,
std::vector<std::string>& dmd_paths,
Vector* desired_point,
std::string rbf = "G",
std::string interp_method = "LS",
double closest_rbf_val = 0.9,
bool reorthogonalize_W = false)
{
T *parametric_dmd = NULL;
getParametricDMD(parametric_dmd, parameter_points, dmd_paths, desired_point,
rbf, interp_method, closest_rbf_val, reorthogonalize_W);
return parametric_dmd;
}

void init_ParametricDMD(pybind11::module_ &m) {

// original getParametricDMD pass a template pointer reference T* &parametric_dmd.
// While it is impossible to bind a template function as itself,
// this first argument T* &parametric_dmd is mainly for determining the type T.
// Here we introduce a dummy argument in place of parametric_dmd.
// This will let python decide which function to use, based on the first argument type.
// We will need variants of this as we bind more DMD classes,
// where dmd_type is the corresponding type.
m.def("getParametricDMD", [](
py::object &dmd_type,
std::vector<Vector*>& parameter_points,
std::vector<DMD*>& dmds,
Vector* desired_point,
std::string rbf = "G",
std::string interp_method = "LS",
double closest_rbf_val = 0.9,
bool reorthogonalize_W = false)
{
std::string name = dmd_type.attr("__name__").cast<std::string>();
if (name == "DMD")
return getParametricDMD_Type<DMD>(parameter_points, dmds, desired_point,
rbf, interp_method, closest_rbf_val, reorthogonalize_W);
else
{
std::string msg = name + " is not a proper libROM DMD class!\n";
throw std::runtime_error(msg.c_str());
}
},
py::arg("dmd_type"),
py::arg("parameter_points"),
py::arg("dmds"),
py::arg("desired_point"),
py::arg("rbf") = "G",
py::arg("interp_method") = "LS",
py::arg("closest_rbf_val") = 0.9,
py::arg("reorthogonalize_W") = false);

// original getParametricDMD pass a template pointer reference T* &parametric_dmd.
// While it is impossible to bind a template function as itself,
// this first argument T* &parametric_dmd is mainly for determining the type T.
// Here we introduce a dummy argument in place of parametric_dmd.
// This will let python decide which function to use, based on the first argument type.
// We will need variants of this as we bind more DMD classes,
// where dmd_type is the corresponding type.
m.def("getParametricDMD", [](
py::object &dmd_type,
std::vector<Vector*>& parameter_points,
std::vector<std::string>& dmd_paths,
Vector* desired_point,
std::string rbf = "G",
std::string interp_method = "LS",
double closest_rbf_val = 0.9,
bool reorthogonalize_W = false)
{
std::string name = dmd_type.attr("__name__").cast<std::string>();
if (name == "DMD")
return getParametricDMD_Type<DMD>(parameter_points, dmd_paths, desired_point,
rbf, interp_method, closest_rbf_val, reorthogonalize_W);
else
{
std::string msg = name + " is not a proper libROM DMD class!\n";
throw std::runtime_error(msg.c_str());
}
},
py::arg("dmd_type"),
py::arg("parameter_points"),
py::arg("dmd_paths"),
py::arg("desired_point"),
py::arg("rbf") = "G",
py::arg("interp_method") = "LS",
py::arg("closest_rbf_val") = 0.9,
py::arg("reorthogonalize_W") = false);

}
94 changes: 94 additions & 0 deletions bindings/pylibROM/mfem/PointwiseSnapshot.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,94 @@
import numpy as np
import mfem.par as mfem
from mpi4py import MPI

class PointwiseSnapshot:
finder = None
npoints = -1
dims = [-1] * 3
spaceDim = -1

domainMin = mfem.Vector()
domainMax = mfem.Vector()
xyz = mfem.Vector()

def __init__(self, sdim, dims_):
self.finder = None
self.spaceDim = sdim
assert((1 < sdim) and (sdim < 4))

self.npoints = np.prod(dims_[:self.spaceDim])
self.dims = np.ones((3,), dtype=int)
self.dims[:self.spaceDim] = dims_[:self.spaceDim]

self.xyz.SetSize(self.npoints * self.spaceDim)
self.xyz.Assign(0.)
return

def SetMesh(self, pmesh):
if (self.finder is not None):
self.finder.FreeData() # Free the internal gslib data.
del self.finder

assert(pmesh.Dimension() == self.spaceDim)
assert(pmesh.SpaceDimension() == self.spaceDim)

self.domainMin, self.domainMax = pmesh.GetBoundingBox(0)

h = [0.] * 3
for i in range(self.spaceDim):
h[i] = (self.domainMax[i] - self.domainMin[i]) / float(self.dims[i] - 1)

rank = pmesh.GetMyRank()
if (rank == 0):
print("PointwiseSnapshot on bounding box from (",
self.domainMin[:self.spaceDim],
") to (", self.domainMax[:self.spaceDim], ")")

# TODO(kevin): we might want to re-write this loop in python manner.
xyzData = self.xyz.GetDataArray()
for k in range(self.dims[2]):
pz = self.domainMin[2] + k * h[2] if (self.spaceDim > 2) else 0.0
osk = k * self.dims[0] * self.dims[1]

for j in range(self.dims[1]):
py = self.domainMin[1] + j * h[1]
osj = (j * self.dims[0]) + osk

for i in range(self.dims[0]):
px = self.domainMin[0] + i * h[0]
xyzData[i + osj] = px
if (self.spaceDim > 1): xyzData[self.npoints + i + osj] = py
if (self.spaceDim > 2): xyzData[2 * self.npoints + i + osj] = pz

self.finder = mfem.FindPointsGSLIB(MPI.COMM_WORLD)
# mfem.FindPointsGSLIB()
self.finder.Setup(pmesh)
self.finder.SetL2AvgType(mfem.FindPointsGSLIB.NONE)
return

def GetSnapshot(self, f, s):
vdim = f.FESpace().GetVDim()
s.SetSize(self.npoints * vdim)

self.finder.Interpolate(self.xyz, f, s)

code_out = self.finder.GetCode()
print(type(code_out))
print(code_out.__dir__())

assert(code_out.Size() == self.npoints)

# Note that Min() and Max() are not defined for Array<unsigned int>
#MFEM_VERIFY(code_out.Min() >= 0 && code_out.Max() < 2, "");
cmin = code_out[0]
cmax = code_out[0]
# TODO(kevin): does this work for mfem array?
for c in code_out:
if (c < cmin):
cmin = c
if (c > cmax):
cmax = c

assert((cmin >= 0) and (cmax < 2))
return
3 changes: 2 additions & 1 deletion bindings/pylibROM/mfem/__init__.py
Original file line number Diff line number Diff line change
@@ -1,4 +1,5 @@
# To add pure python routines to this module,
# either define/import the python routine in this file.
# This will combine both c++ bindings/pure python routines into this module.
from .Utilities import ComputeCtAB
from .Utilities import ComputeCtAB
from .PointwiseSnapshot import PointwiseSnapshot
7 changes: 7 additions & 0 deletions bindings/pylibROM/pylibROM.cpp
Original file line number Diff line number Diff line change
Expand Up @@ -19,10 +19,13 @@ void init_IncrementalSVD(pybind11::module_ &m);

//algo
void init_DMD(pybind11::module_ &);
void init_ParametricDMD(pybind11::module_ &m);

//utils
void init_mpi_utils(pybind11::module_ &m);
void init_Database(pybind11::module_ &m);
void init_HDFDatabase(pybind11::module_ &m);
void init_CSVDatabase(pybind11::module_ &m);

//mfem
// TODO(kevin): We do not bind mfem-related functions until we figure out how to type-cast SWIG Object.
Expand All @@ -33,6 +36,8 @@ PYBIND11_MODULE(_pylibROM, m) {
py::module utils = m.def_submodule("utils");
init_mpi_utils(utils);
init_Database(utils);
init_HDFDatabase(utils);
init_CSVDatabase(utils);

py::module linalg = m.def_submodule("linalg");
init_vector(linalg);
Expand All @@ -42,13 +47,15 @@ PYBIND11_MODULE(_pylibROM, m) {
init_BasisReader(linalg);
init_Options(linalg);
init_NNLSSolver(linalg);

py::module svd = linalg.def_submodule("svd");
init_SVD(svd);
init_StaticSVD(svd);
init_IncrementalSVD(svd);

py::module algo = m.def_submodule("algo");
init_DMD(algo);
init_ParametricDMD(algo);

// py::module mfem = m.def_submodule("mfem");
// init_mfem_Utilities(mfem);
Expand Down
Loading

0 comments on commit b9581d4

Please sign in to comment.